paddlenlp.transformers.bert.modeling 源代码

# Copyright (c) 2020 PaddlePaddle Authors. All Rights Reserved.
#
# Licensed under the Apache License, Version 2.0 (the "License");
# you may not use this file except in compliance with the License.
# You may obtain a copy of the License at
#
#     http://www.apache.org/licenses/LICENSE-2.0
#
# Unless required by applicable law or agreed to in writing, software
# distributed under the License is distributed on an "AS IS" BASIS,
# WITHOUT WARRANTIES OR CONDITIONS OF ANY KIND, either express or implied.
# See the License for the specific language governing permissions and
# limitations under the License.

import paddle
import paddle.nn as nn
import paddle.nn.functional as F
from paddle.nn import Layer

from .. import PretrainedModel, register_base_model

__all__ = [
    'BertModel',
    "BertPretrainedModel",
    'BertForPretraining',
    'BertPretrainingCriterion',
    'BertPretrainingHeads',
    'BertForSequenceClassification',
    'BertForTokenClassification',
    'BertForQuestionAnswering',
    'BertForMultipleChoice',
    "BertForMaskedLM",
]


class BertEmbeddings(Layer):
    """
    Include embeddings from word, position and token_type embeddings
    """

    def __init__(self,
                 vocab_size,
                 hidden_size=768,
                 hidden_dropout_prob=0.1,
                 max_position_embeddings=512,
                 type_vocab_size=16):
        super(BertEmbeddings, self).__init__()
        self.word_embeddings = nn.Embedding(vocab_size, hidden_size)
        self.position_embeddings = nn.Embedding(max_position_embeddings,
                                                hidden_size)
        self.token_type_embeddings = nn.Embedding(type_vocab_size, hidden_size)
        self.layer_norm = nn.LayerNorm(hidden_size)
        self.dropout = nn.Dropout(hidden_dropout_prob)

    def forward(self, input_ids, token_type_ids=None, position_ids=None):
        if position_ids is None:
            ones = paddle.ones_like(input_ids, dtype="int64")
            seq_length = paddle.cumsum(ones, axis=-1)

            position_ids = seq_length - ones
            position_ids.stop_gradient = True
        if token_type_ids is None:
            token_type_ids = paddle.zeros_like(input_ids, dtype="int64")

        input_embedings = self.word_embeddings(input_ids)
        position_embeddings = self.position_embeddings(position_ids)
        token_type_embeddings = self.token_type_embeddings(token_type_ids)

        embeddings = input_embedings + position_embeddings + token_type_embeddings
        embeddings = self.layer_norm(embeddings)
        embeddings = self.dropout(embeddings)
        return embeddings


class BertPooler(Layer):
    """
    Pool the result of BertEncoder.
    """

    def __init__(self, hidden_size, pool_act="tanh"):
        super(BertPooler, self).__init__()
        self.dense = nn.Linear(hidden_size, hidden_size)
        self.activation = nn.Tanh()
        self.pool_act = pool_act

    def forward(self, hidden_states):
        # We "pool" the model by simply taking the hidden state corresponding
        # to the first token.
        first_token_tensor = hidden_states[:, 0]
        pooled_output = self.dense(first_token_tensor)
        if self.pool_act == "tanh":
            pooled_output = self.activation(pooled_output)
        return pooled_output


[文档]class BertPretrainedModel(PretrainedModel): """ An abstract class for pretrained BERT models. It provides BERT related `model_config_file`, `resource_files_names`, `pretrained_resource_files_map`, `pretrained_init_configuration`, `base_model_prefix` for downloading and loading pretrained models. See :class:`~paddlenlp.transformers.model_utils.PretrainedModel` for more details. """ model_config_file = "model_config.json" pretrained_init_configuration = { "bert-base-uncased": { "vocab_size": 30522, "hidden_size": 768, "num_hidden_layers": 12, "num_attention_heads": 12, "intermediate_size": 3072, "hidden_act": "gelu", "hidden_dropout_prob": 0.1, "attention_probs_dropout_prob": 0.1, "max_position_embeddings": 512, "type_vocab_size": 2, "initializer_range": 0.02, "pad_token_id": 0, }, "bert-large-uncased": { "vocab_size": 30522, "hidden_size": 1024, "num_hidden_layers": 24, "num_attention_heads": 16, "intermediate_size": 4096, "hidden_act": "gelu", "hidden_dropout_prob": 0.1, "attention_probs_dropout_prob": 0.1, "max_position_embeddings": 512, "type_vocab_size": 2, "initializer_range": 0.02, "pad_token_id": 0, }, "bert-base-multilingual-uncased": { "vocab_size": 105879, "hidden_size": 768, "num_hidden_layers": 12, "num_attention_heads": 12, "intermediate_size": 3072, "hidden_act": "gelu", "hidden_dropout_prob": 0.1, "attention_probs_dropout_prob": 0.1, "max_position_embeddings": 512, "type_vocab_size": 2, "initializer_range": 0.02, "pad_token_id": 0, }, "bert-base-cased": { "vocab_size": 28996, "hidden_size": 768, "num_hidden_layers": 12, "num_attention_heads": 12, "intermediate_size": 3072, "hidden_act": "gelu", "hidden_dropout_prob": 0.1, "attention_probs_dropout_prob": 0.1, "max_position_embeddings": 512, "type_vocab_size": 2, "initializer_range": 0.02, "pad_token_id": 0, }, "bert-base-chinese": { "vocab_size": 21128, "hidden_size": 768, "num_hidden_layers": 12, "num_attention_heads": 12, "intermediate_size": 3072, "hidden_act": "gelu", "hidden_dropout_prob": 0.1, "attention_probs_dropout_prob": 0.1, "max_position_embeddings": 512, "type_vocab_size": 2, "initializer_range": 0.02, "pad_token_id": 0, }, "bert-base-multilingual-cased": { "vocab_size": 119547, "hidden_size": 768, "num_hidden_layers": 12, "num_attention_heads": 12, "intermediate_size": 3072, "hidden_act": "gelu", "hidden_dropout_prob": 0.1, "attention_probs_dropout_prob": 0.1, "max_position_embeddings": 512, "type_vocab_size": 2, "initializer_range": 0.02, "pad_token_id": 0, }, "bert-large-cased": { "vocab_size": 28996, "hidden_size": 1024, "num_hidden_layers": 24, "num_attention_heads": 16, "intermediate_size": 4096, "hidden_act": "gelu", "hidden_dropout_prob": 0.1, "attention_probs_dropout_prob": 0.1, "max_position_embeddings": 512, "type_vocab_size": 2, "initializer_range": 0.02, "pad_token_id": 0, }, "bert-wwm-chinese": { "vocab_size": 21128, "hidden_size": 768, "num_hidden_layers": 12, "num_attention_heads": 12, "intermediate_size": 3072, "hidden_act": "gelu", "hidden_dropout_prob": 0.1, "attention_probs_dropout_prob": 0.1, "max_position_embeddings": 512, "type_vocab_size": 2, "initializer_range": 0.02, "pad_token_id": 0, }, "bert-wwm-ext-chinese": { "vocab_size": 21128, "hidden_size": 768, "num_hidden_layers": 12, "num_attention_heads": 12, "intermediate_size": 3072, "hidden_act": "gelu", "hidden_dropout_prob": 0.1, "attention_probs_dropout_prob": 0.1, "max_position_embeddings": 512, "type_vocab_size": 2, "initializer_range": 0.02, "pad_token_id": 0, }, "macbert-base-chinese": { "vocab_size": 21128, "hidden_size": 768, "num_hidden_layers": 12, "num_attention_heads": 12, "intermediate_size": 3072, "hidden_act": "gelu", "hidden_dropout_prob": 0.1, "attention_probs_dropout_prob": 0.1, "max_position_embeddings": 512, "type_vocab_size": 2, "initializer_range": 0.02, "pad_token_id": 0, }, "macbert-large-chinese": { "vocab_size": 21128, "hidden_size": 1024, "num_hidden_layers": 24, "num_attention_heads": 16, "intermediate_size": 4096, "hidden_act": "gelu", "hidden_dropout_prob": 0.1, "attention_probs_dropout_prob": 0.1, "max_position_embeddings": 512, "type_vocab_size": 2, "initializer_range": 0.02, "pad_token_id": 0, }, "simbert-base-chinese": { "vocab_size": 13685, "hidden_size": 768, "num_hidden_layers": 12, "num_attention_heads": 12, "intermediate_size": 3072, "hidden_act": "gelu", "hidden_dropout_prob": 0.1, "attention_probs_dropout_prob": 0.1, "max_position_embeddings": 512, "type_vocab_size": 2, "initializer_range": 0.02, "pad_token_id": 0, }, } resource_files_names = {"model_state": "model_state.pdparams"} pretrained_resource_files_map = { "model_state": { "bert-base-uncased": "https://bj.bcebos.com/paddlenlp/models/transformers/bert-base-uncased.pdparams", "bert-large-uncased": "https://bj.bcebos.com/paddlenlp/models/transformers/bert-large-uncased.pdparams", "bert-base-multilingual-uncased": "http://bj.bcebos.com/paddlenlp/models/transformers/bert-base-multilingual-uncased.pdparams", "bert-base-cased": "http://bj.bcebos.com/paddlenlp/models/transformers/bert/bert-base-cased.pdparams", "bert-base-chinese": "http://bj.bcebos.com/paddlenlp/models/transformers/bert/bert-base-chinese.pdparams", "bert-base-multilingual-cased": "http://bj.bcebos.com/paddlenlp/models/transformers/bert/bert-base-multilingual-cased.pdparams", "bert-large-cased": "http://bj.bcebos.com/paddlenlp/models/transformers/bert/bert-large-cased.pdparams", "bert-wwm-chinese": "http://bj.bcebos.com/paddlenlp/models/transformers/bert/bert-wwm-chinese.pdparams", "bert-wwm-ext-chinese": "http://bj.bcebos.com/paddlenlp/models/transformers/bert/bert-wwm-ext-chinese.pdparams", "macbert-base-chinese": "https://bj.bcebos.com/paddlenlp/models/transformers/macbert/macbert-base-chinese.pdparams", "macbert-large-chinese": "https://bj.bcebos.com/paddlenlp/models/transformers/macbert/macbert-large-chinese.pdparams", "simbert-base-chinese": "https://bj.bcebos.com/paddlenlp/models/transformers/simbert/simbert-base-chinese-v1.pdparams", } } base_model_prefix = "bert"
[文档] def init_weights(self, layer): """ Initialization hook """ if isinstance(layer, (nn.Linear, nn.Embedding)): # In the dygraph mode, use the `set_value` to reset the parameter directly, # and reset the `state_dict` to update parameter in static mode. if isinstance(layer.weight, paddle.Tensor): layer.weight.set_value( paddle.tensor.normal( mean=0.0, std=self.initializer_range if hasattr(self, "initializer_range") else self.bert.config["initializer_range"], shape=layer.weight.shape)) elif isinstance(layer, nn.LayerNorm): layer._epsilon = 1e-12
[文档]@register_base_model class BertModel(BertPretrainedModel): """ The bare BERT Model transformer outputting raw hidden-states. This model inherits from :class:`~paddlenlp.transformers.model_utils.PretrainedModel`. Refer to the superclass documentation for the generic methods. This model is also a Paddle `paddle.nn.Layer <https://www.paddlepaddle.org.cn/documentation /docs/en/api/paddle/fluid/dygraph/layers/Layer_en.html>`__ subclass. Use it as a regular Paddle Layer and refer to the Paddle documentation for all matter related to general usage and behavior. Args: vocab_size (int): Vocabulary size of `inputs_ids` in `BertModel`. Also is the vocab size of token embedding matrix. Defines the number of different tokens that can be represented by the `inputs_ids` passed when calling `BertModel`. hidden_size (int, optional): Dimensionality of the embedding layer, encoder layer and pooler layer. Defaults to `768`. num_hidden_layers (int, optional): Number of hidden layers in the Transformer encoder. Defaults to `12`. num_attention_heads (int, optional): Number of attention heads for each attention layer in the Transformer encoder. Defaults to `12`. intermediate_size (int, optional): Dimensionality of the feed-forward (ff) layer in the encoder. Input tensors to ff layers are firstly projected from `hidden_size` to `intermediate_size`, and then projected back to `hidden_size`. Typically `intermediate_size` is larger than `hidden_size`. Defaults to `3072`. hidden_act (str, optional): The non-linear activation function in the feed-forward layer. ``"gelu"``, ``"relu"`` and any other paddle supported activation functions are supported. Defaults to `"gelu"`. hidden_dropout_prob (float, optional): The dropout probability for all fully connected layers in the embeddings and encoder. Defaults to `0.1`. attention_probs_dropout_prob (float, optional): The dropout probability used in MultiHeadAttention in all encoder layers to drop some attention target. Defaults to `0.1`. max_position_embeddings (int, optional): The maximum value of the dimensionality of position encoding, which dictates the maximum supported length of an input sequence. Defaults to `512`. type_vocab_size (int, optional): The vocabulary size of `token_type_ids`. Defaults to `16`. initializer_range (float, optional): The standard deviation of the normal initializer. Defaults to 0.02. .. note:: A normal_initializer initializes weight matrices as normal distributions. See :meth:`BertPretrainedModel.init_weights()` for how weights are initialized in `BertModel`. pad_token_id (int, optional): The index of padding token in the token vocabulary. Defaults to `0`. pooled_act (str, optional): The non-linear activation function in the pooling layer. Defaults to `"tanh"`. """ def __init__(self, vocab_size, hidden_size=768, num_hidden_layers=12, num_attention_heads=12, intermediate_size=3072, hidden_act="gelu", hidden_dropout_prob=0.1, attention_probs_dropout_prob=0.1, max_position_embeddings=512, type_vocab_size=16, initializer_range=0.02, pad_token_id=0, pool_act="tanh"): super(BertModel, self).__init__() self.pad_token_id = pad_token_id self.initializer_range = initializer_range self.embeddings = BertEmbeddings( vocab_size, hidden_size, hidden_dropout_prob, max_position_embeddings, type_vocab_size) encoder_layer = nn.TransformerEncoderLayer( hidden_size, num_attention_heads, intermediate_size, dropout=hidden_dropout_prob, activation=hidden_act, attn_dropout=attention_probs_dropout_prob, act_dropout=0) self.encoder = nn.TransformerEncoder(encoder_layer, num_hidden_layers) self.pooler = BertPooler(hidden_size, pool_act) self.apply(self.init_weights)
[文档] def forward(self, input_ids, token_type_ids=None, position_ids=None, attention_mask=None, output_hidden_states=False): r''' The BertModel forward method, overrides the `__call__()` special method. Args: input_ids (Tensor): Indices of input sequence tokens in the vocabulary. They are numerical representations of tokens that build the input sequence. Its data type should be `int64` and it has a shape of [batch_size, sequence_length]. token_type_ids (Tensor, optional): Segment token indices to indicate different portions of the inputs. Selected in the range ``[0, type_vocab_size - 1]``. If `type_vocab_size` is 2, which means the inputs have two portions. Indices can either be 0 or 1: - 0 corresponds to a *sentence A* token, - 1 corresponds to a *sentence B* token. Its data type should be `int64` and it has a shape of [batch_size, sequence_length]. Defaults to `None`, which means we don't add segment embeddings. position_ids(Tensor, optional): Indices of positions of each input sequence tokens in the position embeddings. Selected in the range ``[0, max_position_embeddings - 1]``. Shape as `(batch_size, num_tokens)` and dtype as int64. Defaults to `None`. attention_mask (Tensor, optional): Mask used in multi-head attention to avoid performing attention on to some unwanted positions, usually the paddings or the subsequent positions. Its data type can be int, float and bool. When the data type is bool, the `masked` tokens have `False` values and the others have `True` values. When the data type is int, the `masked` tokens have `0` values and the others have `1` values. When the data type is float, the `masked` tokens have `-INF` values and the others have `0` values. It is a tensor with shape broadcasted to `[batch_size, num_attention_heads, sequence_length, sequence_length]`. Defaults to `None`, which means nothing needed to be prevented attention to. output_hidden_states (bool, optional): Whether to return the output of each hidden layers. Defaults to `False`. Returns: tuple: Returns tuple (`sequence_output`, `pooled_output`) or (`encoder_outputs`, `pooled_output`). With the fields: - `sequence_output` (Tensor): Sequence of hidden-states at the last layer of the model. It's data type should be float32 and its shape is [batch_size, sequence_length, hidden_size]. - `pooled_output` (Tensor): The output of first token (`[CLS]`) in sequence. We "pool" the model by simply taking the hidden state corresponding to the first token. Its data type should be float32 and its shape is [batch_size, hidden_size]. - `encoder_outputs` (List(Tensor)): A list of Tensor containing hidden-states of the model at each hidden layer in the Transformer encoder. The length of the list is `num_hidden_layers`. Each Tensor has a data type of float32 and its shape is [batch_size, sequence_length, hidden_size]. Example: .. code-block:: import paddle from paddlenlp.transformers import BertModel, BertTokenizer tokenizer = BertTokenizer.from_pretrained('bert-wwm-chinese') model = BertModel.from_pretrained('bert-wwm-chinese') inputs = tokenizer("欢迎使用百度飞桨!") inputs = {k:paddle.to_tensor([v]) for (k, v) in inputs.items()} output = model(**inputs) ''' if attention_mask is None: attention_mask = paddle.unsqueeze( (input_ids == self.pad_token_id ).astype(self.pooler.dense.weight.dtype) * -1e4, axis=[1, 2]) else: if attention_mask.ndim == 2: # attention_mask [batch_size, sequence_length] -> [batch_size, 1, 1, sequence_length] attention_mask = attention_mask.unsqueeze( axis=[1, 2]).astype(paddle.get_default_dtype()) attention_mask = (1.0 - attention_mask) * -1e4 embedding_output = self.embeddings( input_ids=input_ids, position_ids=position_ids, token_type_ids=token_type_ids) if output_hidden_states: output = embedding_output encoder_outputs = [] for mod in self.encoder.layers: output = mod(output, src_mask=attention_mask) encoder_outputs.append(output) if self.encoder.norm is not None: encoder_outputs[-1] = self.encoder.norm(encoder_outputs[-1]) pooled_output = self.pooler(encoder_outputs[-1]) else: sequence_output = self.encoder(embedding_output, attention_mask) pooled_output = self.pooler(sequence_output) if output_hidden_states: return encoder_outputs, pooled_output else: return sequence_output, pooled_output
[文档]class BertForQuestionAnswering(BertPretrainedModel): """ Bert Model with a linear layer on top of the hidden-states output to compute `span_start_logits` and `span_end_logits`, designed for question-answering tasks like SQuAD. Args: bert (:class:`BertModel`): An instance of BertModel. dropout (float, optional): The dropout probability for output of BERT. If None, use the same value as `hidden_dropout_prob` of `BertModel` instance `bert`. Defaults to `None`. """ def __init__(self, bert, dropout=None): super(BertForQuestionAnswering, self).__init__() self.bert = bert # allow bert to be config self.dropout = nn.Dropout(dropout if dropout is not None else self.bert.config["hidden_dropout_prob"]) self.classifier = nn.Linear(self.bert.config["hidden_size"], 2) self.apply(self.init_weights)
[文档] def forward(self, input_ids, token_type_ids=None, position_ids=None, attention_mask=None): r""" The BertForQuestionAnswering forward method, overrides the __call__() special method. Args: input_ids (Tensor): See :class:`BertModel`. token_type_ids (Tensor, optional): See :class:`BertModel`. Returns: tuple: Returns tuple (`start_logits`, `end_logits`). With the fields: - `start_logits` (Tensor): A tensor of the input token classification logits, indicates the start position of the labelled span. Its data type should be float32 and its shape is [batch_size, sequence_length]. - `end_logits` (Tensor): A tensor of the input token classification logits, indicates the end position of the labelled span. Its data type should be float32 and its shape is [batch_size, sequence_length]. Example: .. code-block:: import paddle from paddlenlp.transformers.bert.modeling import BertForQuestionAnswering from paddlenlp.transformers.bert.tokenizer import BertTokenizer tokenizer = BertTokenizer.from_pretrained('bert-base-cased') model = BertForQuestionAnswering.from_pretrained('bert-base-cased') inputs = tokenizer("Welcome to use PaddlePaddle and PaddleNLP!") inputs = {k:paddle.to_tensor([v]) for (k, v) in inputs.items()} outputs = model(**inputs) start_logits = outputs[0] end_logits = outputs[1] """ sequence_output, _ = self.bert( input_ids, token_type_ids=token_type_ids, position_ids=position_ids, attention_mask=attention_mask) logits = self.classifier(sequence_output) logits = paddle.transpose(logits, perm=[2, 0, 1]) start_logits, end_logits = paddle.unstack(x=logits, axis=0) return start_logits, end_logits
[文档]class BertForSequenceClassification(BertPretrainedModel): """ Bert Model with a linear layer on top of the output layer, designed for sequence classification/regression tasks like GLUE tasks. Args: bert (:class:`BertModel`): An instance of BertModel. num_classes (int, optional): The number of classes. Defaults to `2`. dropout (float, optional): The dropout probability for output of BERT. If None, use the same value as `hidden_dropout_prob` of `BertModel` instance `bert`. Defaults to None. """ def __init__(self, bert, num_classes=2, dropout=None): super(BertForSequenceClassification, self).__init__() self.num_classes = num_classes self.bert = bert # allow bert to be config self.dropout = nn.Dropout(dropout if dropout is not None else self.bert.config["hidden_dropout_prob"]) self.classifier = nn.Linear(self.bert.config["hidden_size"], num_classes) self.apply(self.init_weights)
[文档] def forward(self, input_ids, token_type_ids=None, position_ids=None, attention_mask=None): r""" The BertForSequenceClassification forward method, overrides the __call__() special method. Args: input_ids (Tensor): See :class:`BertModel`. token_type_ids (Tensor, optional): See :class:`BertModel`. position_ids(Tensor, optional): See :class:`BertModel`. attention_mask (list, optional): See :class:`BertModel`. Returns: Tensor: Returns tensor `logits`, a tensor of the input text classification logits. Shape as `[batch_size, num_classes]` and dtype as float32. Example: .. code-block:: import paddle from paddlenlp.transformers.bert.modeling import BertForSequenceClassification from paddlenlp.transformers.bert.tokenizer import BertTokenizer tokenizer = BertTokenizer.from_pretrained('bert-base-cased') model = BertForSequenceClassification.from_pretrained('bert-base-cased', num_classes=2) inputs = tokenizer("Welcome to use PaddlePaddle and PaddleNLP!") inputs = {k:paddle.to_tensor([v]) for (k, v) in inputs.items()} logits = model(**inputs) print(logits.shape) # [1, 2] """ _, pooled_output = self.bert( input_ids, token_type_ids=token_type_ids, position_ids=position_ids, attention_mask=attention_mask) pooled_output = self.dropout(pooled_output) logits = self.classifier(pooled_output) return logits
[文档]class BertForTokenClassification(BertPretrainedModel): """ Bert Model with a linear layer on top of the hidden-states output layer, designed for token classification tasks like NER tasks. Args: bert (:class:`BertModel`): An instance of BertModel. num_classes (int, optional): The number of classes. Defaults to `2`. dropout (float, optional): The dropout probability for output of BERT. If None, use the same value as `hidden_dropout_prob` of `BertModel` instance `bert`. Defaults to None. """ def __init__(self, bert, num_classes=2, dropout=None): super(BertForTokenClassification, self).__init__() self.num_classes = num_classes self.bert = bert # allow bert to be config self.dropout = nn.Dropout(dropout if dropout is not None else self.bert.config["hidden_dropout_prob"]) self.classifier = nn.Linear(self.bert.config["hidden_size"], num_classes) self.apply(self.init_weights)
[文档] def forward(self, input_ids, token_type_ids=None, position_ids=None, attention_mask=None): r""" The BertForTokenClassification forward method, overrides the __call__() special method. Args: input_ids (Tensor): See :class:`BertModel`. token_type_ids (Tensor, optional): See :class:`BertModel`. position_ids(Tensor, optional): See :class:`BertModel`. attention_mask (list, optional): See :class:`BertModel`. Returns: Tensor: Returns tensor `logits`, a tensor of the input token classification logits. Shape as `[batch_size, sequence_length, num_classes]` and dtype as `float32`. Example: .. code-block:: import paddle from paddlenlp.transformers.bert.modeling import BertForTokenClassification from paddlenlp.transformers.bert.tokenizer import BertTokenizer tokenizer = BertTokenizer.from_pretrained('bert-base-cased') model = BertForTokenClassification.from_pretrained('bert-base-cased', num_classes=2) inputs = tokenizer("Welcome to use PaddlePaddle and PaddleNLP!") inputs = {k:paddle.to_tensor([v]) for (k, v) in inputs.items()} logits = model(**inputs) print(logits.shape) # [1, 13, 2] """ sequence_output, _ = self.bert( input_ids, token_type_ids=token_type_ids, position_ids=position_ids, attention_mask=attention_mask) sequence_output = self.dropout(sequence_output) logits = self.classifier(sequence_output) return logits
class BertLMPredictionHead(Layer): """ Bert Model with a `language modeling` head on top for CLM fine-tuning. """ def __init__(self, hidden_size, vocab_size, activation, embedding_weights=None): super(BertLMPredictionHead, self).__init__() self.transform = nn.Linear(hidden_size, hidden_size) self.activation = getattr(nn.functional, activation) self.layer_norm = nn.LayerNorm(hidden_size) self.decoder_weight = self.create_parameter( shape=[vocab_size, hidden_size], dtype=self.transform.weight.dtype, is_bias=False) if embedding_weights is None else embedding_weights self.decoder_bias = self.create_parameter( shape=[vocab_size], dtype=self.decoder_weight.dtype, is_bias=True) def forward(self, hidden_states, masked_positions=None): if masked_positions is not None: hidden_states = paddle.reshape(hidden_states, [-1, hidden_states.shape[-1]]) hidden_states = paddle.tensor.gather(hidden_states, masked_positions) # gather masked tokens might be more quick hidden_states = self.transform(hidden_states) hidden_states = self.activation(hidden_states) hidden_states = self.layer_norm(hidden_states) hidden_states = paddle.tensor.matmul( hidden_states, self.decoder_weight, transpose_y=True) + self.decoder_bias return hidden_states
[文档]class BertPretrainingHeads(Layer): """ Perform language modeling task and next sentence classification task. Args: hidden_size (int): See :class:`BertModel`. vocab_size (int): See :class:`BertModel`. activation (str): Activation function used in the language modeling task. embedding_weights (Tensor, optional): Decoding weights used to map hidden_states to logits of the masked token prediction. Its data type should be float32 and its shape is [vocab_size, hidden_size]. Defaults to `None`, which means use the same weights of the embedding layer. """ def __init__(self, hidden_size, vocab_size, activation, embedding_weights=None): super(BertPretrainingHeads, self).__init__() self.predictions = BertLMPredictionHead(hidden_size, vocab_size, activation, embedding_weights) self.seq_relationship = nn.Linear(hidden_size, 2)
[文档] def forward(self, sequence_output, pooled_output, masked_positions=None): """ Args: sequence_output(Tensor): Sequence of hidden-states at the last layer of the model. It's data type should be float32 and its shape is [batch_size, sequence_length, hidden_size]. pooled_output(Tensor): The output of first token (`[CLS]`) in sequence. We "pool" the model by simply taking the hidden state corresponding to the first token. Its data type should be float32 and its shape is [batch_size, hidden_size]. masked_positions(Tensor, optional): A tensor indicates positions to be masked in the position embedding. Its data type should be int64 and its shape is [batch_size, mask_token_num]. `mask_token_num` is the number of masked tokens. It should be no bigger than `sequence_length`. Defaults to `None`, which means we output hidden-states of all tokens in masked token prediction. Returns: tuple: Returns tuple (``prediction_scores``, ``seq_relationship_score``). With the fields: - `prediction_scores` (Tensor): The scores of masked token prediction. Its data type should be float32. If `masked_positions` is None, its shape is [batch_size, sequence_length, vocab_size]. Otherwise, its shape is [batch_size, mask_token_num, vocab_size]. - `seq_relationship_score` (Tensor): The scores of next sentence prediction. Its data type should be float32 and its shape is [batch_size, 2]. """ prediction_scores = self.predictions(sequence_output, masked_positions) seq_relationship_score = self.seq_relationship(pooled_output) return prediction_scores, seq_relationship_score
[文档]class BertForPretraining(BertPretrainedModel): """ Bert Model with pretraining tasks on top. Args: bert (:class:`BertModel`): An instance of :class:`BertModel`. """ def __init__(self, bert): super(BertForPretraining, self).__init__() self.bert = bert self.cls = BertPretrainingHeads( self.bert.config["hidden_size"], self.bert.config["vocab_size"], self.bert.config["hidden_act"], embedding_weights=self.bert.embeddings.word_embeddings.weight) self.apply(self.init_weights)
[文档] def forward(self, input_ids, token_type_ids=None, position_ids=None, attention_mask=None, masked_positions=None): r""" Args: input_ids (Tensor): See :class:`BertModel`. token_type_ids (Tensor, optional): See :class:`BertModel`. position_ids (Tensor, optional): See :class:`BertModel`. attention_mask (Tensor, optional): See :class:`BertModel`. masked_positions(Tensor, optional): See :class:`BertPretrainingHeads`. Returns: tuple: Returns tuple (``prediction_scores``, ``seq_relationship_score``). With the fields: - `prediction_scores` (Tensor): The scores of masked token prediction. Its data type should be float32. If `masked_positions` is None, its shape is [batch_size, sequence_length, vocab_size]. Otherwise, its shape is [batch_size, mask_token_num, vocab_size]. - `seq_relationship_score` (Tensor): The scores of next sentence prediction. Its data type should be float32 and its shape is [batch_size, 2]. """ with paddle.static.amp.fp16_guard(): outputs = self.bert( input_ids, token_type_ids=token_type_ids, position_ids=position_ids, attention_mask=attention_mask) sequence_output, pooled_output = outputs[:2] prediction_scores, seq_relationship_score = self.cls( sequence_output, pooled_output, masked_positions) return prediction_scores, seq_relationship_score
[文档]class BertPretrainingCriterion(paddle.nn.Layer): """ Args: vocab_size(int): Vocabulary size of `inputs_ids` in `BertModel`. Defines the number of different tokens that can be represented by the `inputs_ids` passed when calling `BertModel`. """ def __init__(self, vocab_size): super(BertPretrainingCriterion, self).__init__() # CrossEntropyLoss is expensive since the inner reshape (copy) self.loss_fn = paddle.nn.loss.CrossEntropyLoss(ignore_index=-1) self.vocab_size = vocab_size
[文档] def forward(self, prediction_scores, seq_relationship_score, masked_lm_labels, next_sentence_labels, masked_lm_scale): """ Args: prediction_scores(Tensor): The scores of masked token prediction. Its data type should be float32. If `masked_positions` is None, its shape is [batch_size, sequence_length, vocab_size]. Otherwise, its shape is [batch_size, mask_token_num, vocab_size] seq_relationship_score(Tensor): The scores of next sentence prediction. Its data type should be float32 and its shape is [batch_size, 2] masked_lm_labels(Tensor): The labels of the masked language modeling, its dimensionality is equal to `prediction_scores`. Its data type should be int64. If `masked_positions` is None, its shape is [batch_size, sequence_length, 1]. Otherwise, its shape is [batch_size, mask_token_num, 1] next_sentence_labels(Tensor): The labels of the next sentence prediction task, the dimensionality of `next_sentence_labels` is equal to `seq_relation_labels`. Its data type should be int64 and its shape is [batch_size, 1] masked_lm_scale(Tensor or int): The scale of masked tokens. Used for the normalization of masked language modeling loss. If it is a `Tensor`, its data type should be int64 and its shape is equal to `prediction_scores`. Returns: Tensor: The pretraining loss, equals to the sum of `masked_lm_loss` plus the mean of `next_sentence_loss`. Its data type should be float32 and its shape is [1]. """ with paddle.static.amp.fp16_guard(): masked_lm_loss = F.cross_entropy( prediction_scores, masked_lm_labels, reduction='none', ignore_index=-1) masked_lm_loss = masked_lm_loss / masked_lm_scale next_sentence_loss = F.cross_entropy( seq_relationship_score, next_sentence_labels, reduction='none') return paddle.sum(masked_lm_loss) + paddle.mean(next_sentence_loss)
[文档]class BertForMultipleChoice(BertPretrainedModel): """ Bert Model with a linear layer on top of the hidden-states output layer, designed for multiple choice tasks like RocStories/SWAG tasks. Args: bert (:class:`BertModel`): An instance of BertModel. num_choices (int, optional): The number of choices. Defaults to `2`. dropout (float, optional): The dropout probability for output of Bert. If None, use the same value as `hidden_dropout_prob` of `BertModel` instance `bert`. Defaults to None. """ def __init__(self, bert, num_choices=2, dropout=None): super(BertForMultipleChoice, self).__init__() self.num_choices = num_choices self.bert = bert self.dropout = nn.Dropout(dropout if dropout is not None else self.bert.config["hidden_dropout_prob"]) self.classifier = nn.Linear(self.bert.config["hidden_size"], 1) self.apply(self.init_weights)
[文档] def forward(self, input_ids, token_type_ids=None, position_ids=None, attention_mask=None): r""" The BertForMultipleChoice forward method, overrides the __call__() special method. Args: input_ids (Tensor): See :class:`BertModel` and shape as [batch_size, num_choice, sequence_length]. token_type_ids(Tensor, optional): See :class:`BertModel` and shape as [batch_size, num_choice, sequence_length]. position_ids(Tensor, optional): See :class:`BertModel` and shape as [batch_size, num_choice, sequence_length]. attention_mask (list, optional): See :class:`BertModel` and shape as [batch_size, num_choice, sequence_length]. Returns: Tensor: Returns tensor `reshaped_logits`, a tensor of the multiple choice classification logits. Shape as `[batch_size, num_choice]` and dtype as `float32`. Example: .. code-block:: import paddle from paddlenlp.transformers import BertForMultipleChoice, BertTokenizer from paddlenlp.data import Pad, Dict tokenizer = BertTokenizer.from_pretrained('bert-base-uncased') model = BertForMultipleChoice.from_pretrained('bert-base-uncased', num_choices=2) data = [ { "question": "how do you turn on an ipad screen?", "answer1": "press the volume button.", "answer2": "press the lock button.", "label": 1, }, { "question": "how do you indent something?", "answer1": "leave a space before starting the writing", "answer2": "press the spacebar", "label": 0, }, ] text = [] text_pair = [] for d in data: text.append(d["question"]) text_pair.append(d["answer1"]) text.append(d["question"]) text_pair.append(d["answer2"]) inputs = tokenizer(text, text_pair) batchify_fn = lambda samples, fn=Dict( { "input_ids": Pad(axis=0, pad_val=tokenizer.pad_token_id), # input_ids "token_type_ids": Pad( axis=0, pad_val=tokenizer.pad_token_type_id ), # token_type_ids } ): fn(samples) inputs = batchify_fn(inputs) reshaped_logits = model( input_ids=paddle.to_tensor(inputs[0], dtype="int64"), token_type_ids=paddle.to_tensor(inputs[1], dtype="int64"), ) print(reshaped_logits.shape) # [2, 2] """ # input_ids: [bs, num_choice, seq_l] input_ids = input_ids.reshape(shape=( -1, input_ids.shape[-1])) # flat_input_ids: [bs*num_choice,seq_l] if position_ids is not None: position_ids = position_ids.reshape(shape=(-1, position_ids.shape[-1])) if token_type_ids is not None: token_type_ids = token_type_ids.reshape(shape=( -1, token_type_ids.shape[-1])) if attention_mask is not None: attention_mask = attention_mask.reshape( shape=(-1, attention_mask.shape[-1])) _, pooled_output = self.bert( input_ids, token_type_ids=token_type_ids, position_ids=position_ids, attention_mask=attention_mask) pooled_output = self.dropout(pooled_output) logits = self.classifier(pooled_output) # logits: (bs*num_choice,1) reshaped_logits = logits.reshape( shape=(-1, self.num_choices)) # logits: (bs, num_choice) return reshaped_logits
class BertOnlyMLMHead(nn.Layer): def __init__(self, hidden_size, vocab_size, activation, embedding_weights): super().__init__() self.predictions = BertLMPredictionHead( hidden_size=hidden_size, vocab_size=vocab_size, activation=activation, embedding_weights=embedding_weights) def forward(self, sequence_output, masked_positions=None): prediction_scores = self.predictions(sequence_output, masked_positions) return prediction_scores
[文档]class BertForMaskedLM(BertPretrainedModel): """ Bert Model with a `masked language modeling` head on top. Args: bert (:class:`BertModel`): An instance of :class:`BertModel`. """ def __init__(self, bert): super(BertForMaskedLM, self).__init__() self.bert = bert self.cls = BertOnlyMLMHead( self.bert.config["hidden_size"], self.bert.config["vocab_size"], self.bert.config["hidden_act"], embedding_weights=self.bert.embeddings.word_embeddings.weight) self.apply(self.init_weights)
[文档] def forward(self, input_ids, token_type_ids=None, position_ids=None, attention_mask=None): r""" Args: input_ids (Tensor): See :class:`BertModel`. token_type_ids (Tensor, optional): See :class:`BertModel`. position_ids (Tensor, optional): See :class:`BertModel`. attention_mask (Tensor, optional): See :class:`BertModel`. Returns: Tensor: Returns tensor `prediction_scores`, The scores of masked token prediction. Its data type should be float32 and shape is [batch_size, sequence_length, vocab_size]. Example: .. code-block:: import paddle from paddlenlp.transformers import BertForMaskedLM, BertTokenizer tokenizer = BertTokenizer.from_pretrained('bert-base-uncased') model = BertForMaskedLM.from_pretrained('bert-base-uncased') inputs = tokenizer("Welcome to use PaddlePaddle and PaddleNLP!") inputs = {k:paddle.to_tensor([v]) for (k, v) in inputs.items()} logits = model(**inputs) print(logits.shape) # [1, 13, 30522] """ outputs = self.bert( input_ids, token_type_ids=token_type_ids, position_ids=position_ids, attention_mask=attention_mask) sequence_output = outputs[0] prediction_scores = self.cls(sequence_output, masked_positions=None) return prediction_scores