Transformers documentation
DBRX
DBRX
๊ฐ์
DBRX๋ ํธ๋์คํฌ๋จธ ๊ธฐ๋ฐ์ ๋ค์ ํ ํฐ์ ์์ธกํ๋ ๋์ฝ๋ ์ ์ฉ LLM ๋ชจ๋ธ์ ๋๋ค. ์ด 132B ๋งค๊ฐ๋ณ์๋ฅผ ๊ฐ์ง ์ธ๋ฐํ ์ ๋ฌธ๊ฐ ํผํฉ(MoE) ์ํคํ ์ฒ๋ฅผ ์ฌ์ฉํ๋ฉฐ, ์ด ์ค 36B ๋งค๊ฐ๋ณ์๊ฐ ์ ๋ ฅ๋ง๋ค ํ์ฑํ๋ฉ๋๋ค. 12T ํ ํฐ์ ํ ์คํธ์ ์ฝ๋ ๋ฐ์ดํฐ๋ก ์ฌ์ ํ์ต๋์์ต๋๋ค.
Mixtral-8x7B์ Grok-1๊ณผ ๊ฐ์ ๋ค๋ฅธ ๊ณต๊ฐ MoE ๋ชจ๋ธ๋ค๊ณผ ๋น๊ตํ์ ๋, DBRX๋ ๋ ๋ง์ ์์ ์์ ์ ๋ฌธ๊ฐ๋ค์ ์ฌ์ฉํ๋ ์ธ๋ฐํ ๊ตฌ์กฐ๋ฅผ ๊ฐ์ง๊ณ ์์ต๋๋ค. DBRX๋ 16๊ฐ์ ์ ๋ฌธ๊ฐ ์ค 4๊ฐ๋ฅผ ์ ํํ๋ ๋ฐ๋ฉด, Mixtral-8x7B์ Grok-1์ 8๊ฐ์ ์ ๋ฌธ๊ฐ ์ค 2๊ฐ๋ฅผ ์ ํํฉ๋๋ค.
์ด๋ 65๋ฐฐ ๋ ๋ง์ ์ ๋ฌธ๊ฐ ์กฐํฉ์ ๊ฐ๋ฅํ๊ฒ ํ๋ฉฐ, ์ด๋ฅผ ํตํด ๋ชจ๋ธ์ ํ์ง์ด ํฅ์๋๋ ๊ฒ์ ๋ฐ๊ฒฌํ์ต๋๋ค. DBRX๋ ํ์ ์์น ์ธ์ฝ๋ฉ(RoPE), ๊ฒ์ดํธ ์ ํ ์ ๋(GLU), ๊ทธ๋ฃน ์ฟผ๋ฆฌ ์ดํ ์ (GQA)์ ์ฌ์ฉํฉ๋๋ค. BPE ๊ธฐ๋ฐ ๋ชจ๋ธ์ด๋ฉฐ tiktoken ์ ์ฅ์์ ์ค๋ช ๋ GPT-4 ํ ํฌ๋์ด์ ๋ฅผ ์ฌ์ฉํฉ๋๋ค. ์ด๋ฌํ ์ ํ๋ค์ ์ฒ ์ ํ ํ๊ฐ์ ์ค์ผ์ผ๋ง ์คํ์ ๊ธฐ๋ฐ์ผ๋ก ์ด๋ฃจ์ด์ก์ต๋๋ค.
DBRX๋ ์ ์คํ๊ฒ ์ ๋ณ๋ 12T ํ ํฐ์ ๋ฐ์ดํฐ๋ก ์ฌ์ ํ์ต๋์์ผ๋ฉฐ, ์ต๋ ๋ฌธ๋งฅ ๊ธธ์ด๋ 32K ํ ํฐ์ ๋๋ค. ์ด ๋ฐ์ดํฐ๋ ํ ํฐ ๋๋น MPT ๊ณ์ด ๋ชจ๋ธ ํ์ต์ ์ฌ์ฉ๋ ๋ฐ์ดํฐ๋ณด๋ค ์ต์ 2๋ฐฐ ์ด์ ๋ ์ข์ ๊ฒ์ผ๋ก ์ถ์ ๋ฉ๋๋ค. ์ด ์๋ก์ด ๋ฐ์ดํฐ์ ์ ๋ฐ์ดํฐ ์ฒ๋ฆฌ๋ฅผ ์ํ Apache Sparkโข์ Databricks ๋ ธํธ๋ถ, ๊ทธ๋ฆฌ๊ณ ๋ฐ์ดํฐ ๊ด๋ฆฌ์ ๊ฑฐ๋ฒ๋์ค๋ฅผ ์ํ Unity Catalog๋ฅผ ํฌํจํ Databricks ๋๊ตฌ ์ ์ฒด๋ฅผ ํ์ฉํ์ฌ ๊ฐ๋ฐ๋์์ต๋๋ค. ์ฐ๋ฆฌ๋ ์ฌ์ ํ์ต์ ์ํด ์ปค๋ฆฌํ๋ผ ํ์ต์ ์ฌ์ฉํ์ผ๋ฉฐ, ํ์ต ์ค ๋ฐ์ดํฐ ๋ฏน์ค๋ฅผ ๋ณ๊ฒฝํ๋ ๋ฐฉ์์ด ๋ชจ๋ธ ํ์ง์ ์๋นํ ๊ฐ์ ํ๋ค๋ ๊ฒ์ ๋ฐ๊ฒฌํ์ต๋๋ค.
DBRX Instruct์ DBRX Base์ ๋ํ ๋ ์์ธํ ์ ๋ณด๋ ์ด ๊ธฐ์ ๋ธ๋ก๊ทธ ํฌ์คํธ์์ ํ์ธํ ์ ์์ต๋๋ค.
์ด ๋ชจ๋ธ์ eitan-turok์ abhi-db๊ฐ ๊ธฐ์ฌํ์ต๋๋ค. ์๋ณธ ์ฝ๋๋ ์ด๊ณณ์์ ์ฐพ์ ์ ์์ง๋ง, ์ต์ ๋ฒ์ ์ด ์๋ ์ ์์ต๋๋ค.
์ฌ์ฉ ์
generate()
๋ฉ์๋๋ DBRX๋ฅผ ์ฌ์ฉํ์ฌ ํ
์คํธ๋ฅผ ์์ฑํ๋ ๋ฐ ์ฌ์ฉ๋ ์ ์์ต๋๋ค. ํ์ค ์ดํ
์
๊ตฌํ, ํ๋์ ์ดํ
์
, PyTorch์ ์ค์ผ์ผ๋ ๋ด์ ์ดํ
์
(Scaled Dot-Product Attention)์ ์ฌ์ฉํ์ฌ ์์ฑํ ์ ์์ต๋๋ค. ํ์์ ๋ ์ดํ
์
๊ตฌํ ๋ฐฉ์์ ์ฒ๋ฆฌ ์๋๋ฅผ ํฌ๊ฒ ๋์ฌ์ค๋๋ค.
from transformers import DbrxForCausalLM, AutoTokenizer
import torch
tokenizer = AutoTokenizer.from_pretrained("databricks/dbrx-instruct", token="YOUR_HF_TOKEN")
model = DbrxForCausalLM.from_pretrained(
"databricks/dbrx-instruct",
device_map="auto",
dtype=torch.bfloat16,
token="YOUR_HF_TOKEN",
)
input_text = "What does it take to build a great LLM?"
messages = [{"role": "user", "content": input_text}]
input_ids = tokenizer.apply_chat_template(messages, return_dict=True, tokenize=True, add_generation_prompt=True, return_tensors="pt").to("cuda")
outputs = model.generate(**input_ids, max_new_tokens=200)
print(tokenizer.decode(outputs[0]))
pip install flash-attn
๋ฅผ ํตํด ํ๋์ ์ดํ
์
์ ์ค์นํ๋ฉด, ๋ ๋น ๋ฅธ ์์ฑ์ด ๊ฐ๋ฅํฉ๋๋ค. (ํ๋์ ์ดํ
์
์ ๋ํ HuggingFace ๋ฌธ์๋ ์ด๊ณณ์์ ํ์ธํ ์ ์์ต๋๋ค.)
from transformers import DbrxForCausalLM, AutoTokenizer
import torch
tokenizer = AutoTokenizer.from_pretrained("databricks/dbrx-instruct", token="YOUR_HF_TOKEN")
model = DbrxForCausalLM.from_pretrained(
"databricks/dbrx-instruct",
device_map="auto",
dtype=torch.bfloat16,
token="YOUR_HF_TOKEN",
attn_implementation="flash_attention_2",
)
input_text = "What does it take to build a great LLM?"
messages = [{"role": "user", "content": input_text}]
input_ids = tokenizer.apply_chat_template(messages, return_dict=True, tokenize=True, add_generation_prompt=True, return_tensors="pt").to("cuda")
outputs = model.generate(**input_ids, max_new_tokens=200)
print(tokenizer.decode(outputs[0]))
PyTorch์ ์ค์ผ์ผ๋ ๋ด์ ์ดํ ์ ์ ์ฌ์ฉํ์ฌ๋ ๋ ๋น ๋ฅธ ์์ฑ์ด ๊ฐ๋ฅํฉ๋๋ค. (์ค์ผ์ผ๋ ๋ด์ ์ดํ ์ ์ ๋ํ HuggingFace ๋ฌธ์๋ ์ด๊ณณ์์ ํ์ธํ ์ ์์ต๋๋ค.)
from transformers import DbrxForCausalLM, AutoTokenizer
import torch
tokenizer = AutoTokenizer.from_pretrained("databricks/dbrx-instruct", token="YOUR_HF_TOKEN")
model = DbrxForCausalLM.from_pretrained(
"databricks/dbrx-instruct",
device_map="auto",
dtype=torch.bfloat16,
token="YOUR_HF_TOKEN",
attn_implementation="sdpa",
)
input_text = "What does it take to build a great LLM?"
messages = [{"role": "user", "content": input_text}]
input_ids = tokenizer.apply_chat_template(messages, return_dict=True, tokenize=True, add_generation_prompt=True, return_tensors="pt").to("cuda")
outputs = model.generate(**input_ids, max_new_tokens=200)
print(tokenizer.decode(outputs[0]))
DbrxConfig
class transformers.DbrxConfig
< source >( d_model: int = 2048 n_heads: int = 16 n_layers: int = 24 max_seq_len: int = 2048 vocab_size: int = 32000 resid_pdrop: float = 0.0 emb_pdrop: float = 0.0 attn_config: typing.Optional[transformers.models.dbrx.configuration_dbrx.DbrxAttentionConfig] = None ffn_config: typing.Optional[transformers.models.dbrx.configuration_dbrx.DbrxFFNConfig] = None use_cache: bool = True initializer_range: float = 0.02 output_router_logits: bool = False **kwargs: typing.Any )
Parameters
- d_model (
int
, optional, defaults to 2048) — Dimensionality of the embeddings and hidden states. - n_heads (
int
, optional, defaults to 16) — Number of attention heads for each attention layer in the Transformer encoder. - n_layers (
int
, optional, defaults to 24) — Number of hidden layers in the Transformer encoder. - max_seq_len (
int
, optional, defaults to 2048) — The maximum sequence length of the model. - vocab_size (
int
, optional, defaults to 32000) — Vocabulary size of the Dbrx model. Defines the maximum number of different tokens that can be represented by theinputs_ids
passed when calling DbrxModel. - resid_pdrop (
float
, optional, defaults to 0.0) — The dropout probability applied to the attention output before combining with residual. - emb_pdrop (
float
, optional, defaults to 0.0) — The dropout probability for the embedding layer. - attn_config (
dict
, optional) — A dictionary used to configure the model’s attention module. - ffn_config (
dict
, optional) — A dictionary used to configure the model’s FFN module. - use_cache (
bool
, optional, defaults toTrue
) — Whether or not the model should return the last key/values attentions (not used by all models). - initializer_range (
float
, optional, defaults to 0.02) — The standard deviation of the truncated_normal_initializer for initializing all weight matrices. - output_router_logits (
bool
, optional, defaults toFalse
) — Whether or not the router logits should be returned by the model. Enabling this will also allow the model to output the auxiliary loss. See here for more details.
This is the configuration class to store the configuration of a DbrxModel. It is used to instantiate a Dbrx model according to the specified arguments, defining the model architecture. Instantiating a configuration with the defaults will yield a different configuration to that of the databricks/dbrx-instruct architecture.
Configuration objects inherit from PretrainedConfig and can be used to control the model outputs. Read the documentation from PretrainedConfig for more information.
Example:
>>> from transformers import DbrxConfig, DbrxModel
>>> # Initializing a Dbrx configuration
>>> configuration = DbrxConfig(n_layers=2, d_model=256, n_heads=8, vocab_size=128)
>>> # Initializing a model (with random weights) from the configuration
>>> model = DbrxModel(configuration)
>>> # Accessing the model configuration
>>> configuration = model.config
DbrxModel
class transformers.DbrxModel
< source >( config: DbrxConfig )
Parameters
- config (DbrxConfig) — Model configuration class with all the parameters of the model. Initializing with a config file does not load the weights associated with the model, only the configuration. Check out the from_pretrained() method to load the model weights.
The bare Dbrx Model outputting raw hidden-states without any specific head on top.
This model inherits from PreTrainedModel. Check the superclass documentation for the generic methods the library implements for all its model (such as downloading or saving, resizing the input embeddings, pruning heads etc.)
This model is also a PyTorch torch.nn.Module subclass. Use it as a regular PyTorch Module and refer to the PyTorch documentation for all matter related to general usage and behavior.
forward
< source >( input_ids: typing.Optional[torch.LongTensor] = None attention_mask: typing.Optional[torch.Tensor] = None position_ids: typing.Optional[torch.LongTensor] = None past_key_values: typing.Optional[transformers.cache_utils.Cache] = None inputs_embeds: typing.Optional[torch.Tensor] = None use_cache: typing.Optional[bool] = None output_attentions: typing.Optional[bool] = None output_hidden_states: typing.Optional[bool] = None output_router_logits: typing.Optional[bool] = None return_dict: typing.Optional[bool] = None cache_position: typing.Optional[torch.LongTensor] = None **kwargs ) โ transformers.modeling_outputs.MoeModelOutputWithPast
or tuple(torch.FloatTensor)
Parameters
- input_ids (
torch.LongTensor
of shape(batch_size, sequence_length)
, optional) — Indices of input sequence tokens in the vocabulary. Padding will be ignored by default.Indices can be obtained using AutoTokenizer. See PreTrainedTokenizer.encode() and PreTrainedTokenizer.call() for details.
- attention_mask (
torch.Tensor
of shape(batch_size, sequence_length)
, optional) — Mask to avoid performing attention on padding token indices. Mask values selected in[0, 1]
:- 1 for tokens that are not masked,
- 0 for tokens that are masked.
- position_ids (
torch.LongTensor
of shape(batch_size, sequence_length)
, optional) — Indices of positions of each input sequence tokens in the position embeddings. Selected in the range[0, config.n_positions - 1]
. - past_key_values (
~cache_utils.Cache
, optional) — Pre-computed hidden-states (key and values in the self-attention blocks and in the cross-attention blocks) that can be used to speed up sequential decoding. This typically consists in thepast_key_values
returned by the model at a previous stage of decoding, whenuse_cache=True
orconfig.use_cache=True
.Only Cache instance is allowed as input, see our kv cache guide. If no
past_key_values
are passed, DynamicCache will be initialized by default.The model will output the same cache format that is fed as input.
If
past_key_values
are used, the user is expected to input only unprocessedinput_ids
(those that don’t have their past key value states given to this model) of shape(batch_size, unprocessed_length)
instead of allinput_ids
of shape(batch_size, sequence_length)
. - inputs_embeds (
torch.Tensor
of shape(batch_size, sequence_length, hidden_size)
, optional) — Optionally, instead of passinginput_ids
you can choose to directly pass an embedded representation. This is useful if you want more control over how to convertinput_ids
indices into associated vectors than the model’s internal embedding lookup matrix. - use_cache (
bool
, optional) — If set toTrue
,past_key_values
key value states are returned and can be used to speed up decoding (seepast_key_values
). - output_attentions (
bool
, optional) — Whether or not to return the attentions tensors of all attention layers. Seeattentions
under returned tensors for more detail. - output_hidden_states (
bool
, optional) — Whether or not to return the hidden states of all layers. Seehidden_states
under returned tensors for more detail. - output_router_logits (
bool
, optional) — Whether or not to return the logits of all the routers. They are useful for computing the router loss, and should not be returned during inference. - return_dict (
bool
, optional) — Whether or not to return a ModelOutput instead of a plain tuple. - cache_position (
torch.LongTensor
of shape(sequence_length)
, optional) — Indices depicting the position of the input sequence tokens in the sequence. Contrarily toposition_ids
, this tensor is not affected by padding. It is used to update the cache in the correct position and to infer the complete sequence length.
Returns
transformers.modeling_outputs.MoeModelOutputWithPast
or tuple(torch.FloatTensor)
A transformers.modeling_outputs.MoeModelOutputWithPast
or a tuple of
torch.FloatTensor
(if return_dict=False
is passed or when config.return_dict=False
) comprising various
elements depending on the configuration (DbrxConfig) and inputs.
-
last_hidden_state (
torch.FloatTensor
of shape(batch_size, sequence_length, hidden_size)
) โ Sequence of hidden-states at the output of the last layer of the model. -
past_key_values (
Cache
, optional, returned whenuse_cache=True
is passed or whenconfig.use_cache=True
) โ It is a Cache instance. For more details, see our kv cache guide.Contains pre-computed hidden-states (key and values in the self-attention blocks and optionally if
config.is_encoder_decoder=True
in the cross-attention blocks) that can be used (seepast_key_values
input) to speed up sequential decoding. -
hidden_states (
tuple(torch.FloatTensor)
, optional, returned whenoutput_hidden_states=True
is passed or whenconfig.output_hidden_states=True
) โ Tuple oftorch.FloatTensor
(one for the output of the embeddings, if the model has an embedding layer, + one for the output of each layer) of shape(batch_size, sequence_length, hidden_size)
.Hidden-states of the model at the output of each layer plus the optional initial embedding outputs.
-
attentions (
tuple(torch.FloatTensor)
, optional, returned whenoutput_attentions=True
is passed or whenconfig.output_attentions=True
) โ Tuple oftorch.FloatTensor
(one for each layer) of shape(batch_size, num_heads, sequence_length, sequence_length)
.Attentions weights after the attention softmax, used to compute the weighted average in the self-attention heads.
-
router_logits (
tuple(torch.FloatTensor)
, optional, returned whenoutput_router_probs=True
andconfig.add_router_probs=True
is passed or whenconfig.output_router_probs=True
) โ Tuple oftorch.FloatTensor
(one for each layer) of shape(batch_size, sequence_length, num_experts)
.Raw router logtis (post-softmax) that are computed by MoE routers, these terms are used to compute the auxiliary loss for Mixture of Experts models.
The DbrxModel forward method, overrides the __call__
special method.
Although the recipe for forward pass needs to be defined within this function, one should call the Module
instance afterwards instead of this since the former takes care of running the pre and post processing steps while
the latter silently ignores them.
DbrxForCausalLM
class transformers.DbrxForCausalLM
< source >( config: DbrxConfig )
Parameters
- config (DbrxConfig) — Model configuration class with all the parameters of the model. Initializing with a config file does not load the weights associated with the model, only the configuration. Check out the from_pretrained() method to load the model weights.
The DBRX Model transformer for causal language modeling.
This model inherits from PreTrainedModel. Check the superclass documentation for the generic methods the library implements for all its model (such as downloading or saving, resizing the input embeddings, pruning heads etc.)
This model is also a PyTorch torch.nn.Module subclass. Use it as a regular PyTorch Module and refer to the PyTorch documentation for all matter related to general usage and behavior.
forward
< source >( input_ids: typing.Optional[torch.LongTensor] = None attention_mask: typing.Optional[torch.Tensor] = None position_ids: typing.Optional[torch.LongTensor] = None past_key_values: typing.Optional[transformers.cache_utils.Cache] = None inputs_embeds: typing.Optional[torch.Tensor] = None labels: typing.Optional[torch.LongTensor] = None use_cache: typing.Optional[bool] = None output_attentions: typing.Optional[bool] = None output_hidden_states: typing.Optional[bool] = None output_router_logits: typing.Optional[bool] = None return_dict: typing.Optional[bool] = None cache_position: typing.Optional[torch.LongTensor] = None logits_to_keep: typing.Union[int, torch.Tensor] = 0 **kwargs ) โ transformers.modeling_outputs.MoeCausalLMOutputWithPast
or tuple(torch.FloatTensor)
Parameters
- input_ids (
torch.LongTensor
of shape(batch_size, sequence_length)
, optional) — Indices of input sequence tokens in the vocabulary. Padding will be ignored by default.Indices can be obtained using AutoTokenizer. See PreTrainedTokenizer.encode() and PreTrainedTokenizer.call() for details.
- attention_mask (
torch.Tensor
of shape(batch_size, sequence_length)
, optional) — Mask to avoid performing attention on padding token indices. Mask values selected in[0, 1]
:- 1 for tokens that are not masked,
- 0 for tokens that are masked.
- position_ids (
torch.LongTensor
of shape(batch_size, sequence_length)
, optional) — Indices of positions of each input sequence tokens in the position embeddings. Selected in the range[0, config.n_positions - 1]
. - past_key_values (
~cache_utils.Cache
, optional) — Pre-computed hidden-states (key and values in the self-attention blocks and in the cross-attention blocks) that can be used to speed up sequential decoding. This typically consists in thepast_key_values
returned by the model at a previous stage of decoding, whenuse_cache=True
orconfig.use_cache=True
.Only Cache instance is allowed as input, see our kv cache guide. If no
past_key_values
are passed, DynamicCache will be initialized by default.The model will output the same cache format that is fed as input.
If
past_key_values
are used, the user is expected to input only unprocessedinput_ids
(those that don’t have their past key value states given to this model) of shape(batch_size, unprocessed_length)
instead of allinput_ids
of shape(batch_size, sequence_length)
. - inputs_embeds (
torch.Tensor
of shape(batch_size, sequence_length, hidden_size)
, optional) — Optionally, instead of passinginput_ids
you can choose to directly pass an embedded representation. This is useful if you want more control over how to convertinput_ids
indices into associated vectors than the model’s internal embedding lookup matrix. - labels (
torch.LongTensor
of shape(batch_size, sequence_length)
, optional) — Labels for computing the masked language modeling loss. Indices should either be in[0, ..., config.vocab_size]
or -100 (seeinput_ids
docstring). Tokens with indices set to-100
are ignored (masked), the loss is only computed for the tokens with labels in[0, ..., config.vocab_size]
. - use_cache (
bool
, optional) — If set toTrue
,past_key_values
key value states are returned and can be used to speed up decoding (seepast_key_values
). - output_attentions (
bool
, optional) — Whether or not to return the attentions tensors of all attention layers. Seeattentions
under returned tensors for more detail. - output_hidden_states (
bool
, optional) — Whether or not to return the hidden states of all layers. Seehidden_states
under returned tensors for more detail. - output_router_logits (
bool
, optional) — Whether or not to return the logits of all the routers. They are useful for computing the router loss, and should not be returned during inference. - return_dict (
bool
, optional) — Whether or not to return a ModelOutput instead of a plain tuple. - cache_position (
torch.LongTensor
of shape(sequence_length)
, optional) — Indices depicting the position of the input sequence tokens in the sequence. Contrarily toposition_ids
, this tensor is not affected by padding. It is used to update the cache in the correct position and to infer the complete sequence length. - logits_to_keep (
Union[int, torch.Tensor]
, defaults to0
) — If anint
, compute logits for the lastlogits_to_keep
tokens. If0
, calculate logits for allinput_ids
(special case). Only last token logits are needed for generation, and calculating them only for that token can save memory, which becomes pretty significant for long sequences or large vocabulary size. If atorch.Tensor
, must be 1D corresponding to the indices to keep in the sequence length dimension. This is useful when using packed tensor format (single dimension for batch and sequence length).
Returns
transformers.modeling_outputs.MoeCausalLMOutputWithPast
or tuple(torch.FloatTensor)
A transformers.modeling_outputs.MoeCausalLMOutputWithPast
or a tuple of
torch.FloatTensor
(if return_dict=False
is passed or when config.return_dict=False
) comprising various
elements depending on the configuration (DbrxConfig) and inputs.
-
loss (
torch.FloatTensor
of shape(1,)
, optional, returned whenlabels
is provided) โ Language modeling loss (for next-token prediction). -
logits (
torch.FloatTensor
of shape(batch_size, sequence_length, config.vocab_size)
) โ Prediction scores of the language modeling head (scores for each vocabulary token before SoftMax). -
aux_loss (
torch.FloatTensor
, optional, returned whenlabels
is provided) โ aux_loss for the sparse modules. -
router_logits (
tuple(torch.FloatTensor)
, optional, returned whenoutput_router_probs=True
andconfig.add_router_probs=True
is passed or whenconfig.output_router_probs=True
) โ Tuple oftorch.FloatTensor
(one for each layer) of shape(batch_size, sequence_length, num_experts)
.Raw router logtis (post-softmax) that are computed by MoE routers, these terms are used to compute the auxiliary loss for Mixture of Experts models.
-
past_key_values (
Cache
, optional, returned whenuse_cache=True
is passed or whenconfig.use_cache=True
) โ It is a Cache instance. For more details, see our kv cache guide.Contains pre-computed hidden-states (key and values in the self-attention blocks) that can be used (see
past_key_values
input) to speed up sequential decoding. -
hidden_states (
tuple(torch.FloatTensor)
, optional, returned whenoutput_hidden_states=True
is passed or whenconfig.output_hidden_states=True
) โ Tuple oftorch.FloatTensor
(one for the output of the embeddings, if the model has an embedding layer, + one for the output of each layer) of shape(batch_size, sequence_length, hidden_size)
.Hidden-states of the model at the output of each layer plus the optional initial embedding outputs.
-
attentions (
tuple(torch.FloatTensor)
, optional, returned whenoutput_attentions=True
is passed or whenconfig.output_attentions=True
) โ Tuple oftorch.FloatTensor
(one for each layer) of shape(batch_size, num_heads, sequence_length, sequence_length)
.Attentions weights after the attention softmax, used to compute the weighted average in the self-attention heads.
The DbrxForCausalLM forward method, overrides the __call__
special method.
Although the recipe for forward pass needs to be defined within this function, one should call the Module
instance afterwards instead of this since the former takes care of running the pre and post processing steps while
the latter silently ignores them.
Example:
>> from transformers import AutoTokenizer, DbrxForCausalLM
>> model = DbrxForCausalLM.from_pretrained("databricks/dbrx-instruct")
>> tokenizer = AutoTokenizer.from_pretrained("databricks/dbrx-instruct")
>> prompt = "Hey, are you conscious? Can you talk to me?"
>> inputs = tokenizer(prompt, return_tensors="pt")
>> # Generate
>> generate_ids = model.generate(inputs.input_ids, max_length=30)
>> tokenizer.batch_decode(generate_ids, skip_special_tokens=True, clean_up_tokenization_spaces=False)[0]
"Hey, are you conscious? Can you talk to me?\nI'm not conscious, but I can talk to you."