#![allow(clippy::cast_possible_truncation, clippy::cast_precision_loss)]
use candle_core::{DType, Device, Module, Result, Tensor};
use candle_nn::{RotaryEmbedding, VarBuilder};
use mistralrs_quant::{QuantMethod, QuantMethodConfig, QuantizedConfig, UnquantLinear};
use serde::{Deserialize, Serialize};
use std::{collections::HashMap, sync::Arc};
use crate::{
amoe::AnyMoeBaseModelMixin,
attention::SdpaParams,
device_map::DeviceMapper,
layers::{Activation, CausalMasker, MatMul, RmsNorm, Sdpa},
layers_masker::PastKvLenCache,
paged_attention::{AttentionImplementation, ModelConfigMetadata, PagedAttention},
pipeline::{
extract_logits,
text_models_inputs_processor::{FlashParams, PagedAttentionInputMetadata},
EitherCache, IsqModel, KvCache, NormalCache, NormalLoadingMetadata, NormalModel,
},
serde_default_fn,
utils::{progress::NiceProgressBar, unvarbuilder::UnVarBuilder},
};
serde_default_fn!(bool, word_emb_default, false);
#[derive(Debug, Clone, Deserialize, Serialize)]
pub struct Config {
pub(crate) vocab_size: usize,
pub(crate) hidden_size: usize,
pub(crate) intermediate_size: usize,
pub(crate) num_hidden_layers: usize,
pub(crate) num_attention_heads: usize,
pub(crate) num_key_value_heads: usize,
pub(crate) hidden_act: Activation,
pub(crate) max_position_embeddings: usize,
pub(crate) rms_norm_eps: f64,
pub(crate) rope_theta: f64,
pub(crate) sliding_window: Option<usize>,
pub(crate) num_experts_per_tok: usize,
pub(crate) num_local_experts: usize,
pub(crate) use_flash_attn: bool,
pub(crate) quantization_config: Option<QuantizedConfig>,
#[serde(default = "word_emb_default")]
pub(crate) tie_word_embeddings: bool,
}
struct Attention {
q_proj: Arc<dyn QuantMethod>,
k_proj: Arc<dyn QuantMethod>,
v_proj: Arc<dyn QuantMethod>,
o_proj: Arc<dyn QuantMethod>,
num_heads: usize,
num_kv_heads: usize,
head_dim: usize,
rotary_emb: Arc<RotaryEmbedding>,
sliding_window: Option<usize>,
paged_attn: Option<PagedAttention>,
sdpa_params: SdpaParams,
}
impl Attention {
fn new(
rotary_emb: Arc<RotaryEmbedding>,
cfg: &Config,
vb: VarBuilder,
paged_attn: Option<PagedAttention>,
) -> Result<Self> {
let hidden_sz = cfg.hidden_size;
let num_heads = cfg.num_attention_heads;
let num_kv_heads = cfg.num_key_value_heads;
let head_dim = hidden_sz / num_heads;
let q_proj = mistralrs_quant::linear_no_bias(
hidden_sz,
num_heads * head_dim,
&cfg.quantization_config,
vb.pp("q_proj"),
)?;
let k_proj = mistralrs_quant::linear_no_bias(
hidden_sz,
num_kv_heads * head_dim,
&cfg.quantization_config,
vb.pp("k_proj"),
)?;
let v_proj = mistralrs_quant::linear_no_bias(
hidden_sz,
num_kv_heads * head_dim,
&cfg.quantization_config,
vb.pp("v_proj"),
)?;
let o_proj = mistralrs_quant::linear_no_bias(
num_heads * head_dim,
hidden_sz,
&cfg.quantization_config,
vb.pp("o_proj"),
)?;
Ok(Self {
q_proj,
k_proj,
v_proj,
o_proj,
num_heads,
num_kv_heads,
head_dim,
rotary_emb,
sliding_window: cfg.sliding_window,
paged_attn,
sdpa_params: SdpaParams {
n_kv_groups: num_heads / num_kv_heads,
use_flash_attn: cfg.use_flash_attn,
softcap: None,
softmax_scale: 1.0 / (head_dim as f32).sqrt(),
sliding_window: cfg.sliding_window,
},
})
}
#[allow(clippy::too_many_arguments)]
fn forward(
&self,
xs: &Tensor,
attention_mask: Option<&Tensor>,
seqlen_offsets: &[usize],
start_offsets_kernel: Tensor,
kv_cache: &mut KvCache,
metadata: Option<((Tensor, Tensor), &mut PagedAttentionInputMetadata)>,
flash_params: &FlashParams,
) -> Result<Tensor> {
let (b_sz, q_len, _) = xs.dims3()?;
let original_dtype = xs.dtype();
let mut xs = xs.clone();
if let Some(t) = self.q_proj.quantized_act_type() {
xs = xs.to_dtype(t)?;
}
let mut q = MatMul.qmethod_matmul(&xs, &*self.q_proj)?;
let mut k = MatMul.qmethod_matmul(&xs, &*self.k_proj)?;
let mut v = MatMul.qmethod_matmul(&xs, &*self.v_proj)?;
if self.q_proj.quantized_act_type().is_some() {
q = q.to_dtype(original_dtype)?;
k = k.to_dtype(original_dtype)?;
v = v.to_dtype(original_dtype)?;
}
let mut q = q.reshape((b_sz * q_len, self.num_heads, self.head_dim))?;
let mut k = k.reshape((b_sz * q_len, self.num_kv_heads, self.head_dim))?;
let v = if q_len != 1 {
v.reshape((b_sz, q_len, self.num_kv_heads, self.head_dim))?
.transpose(1, 2)?
} else {
v.reshape((b_sz, self.num_kv_heads, q_len, self.head_dim))?
};
self.rotary_emb
.forward(seqlen_offsets, &start_offsets_kernel, &mut q, &mut k, b_sz)?;
if q.rank() == 3 && q_len != 1 {
q = q
.reshape((b_sz, q_len, self.num_heads, self.head_dim))?
.transpose(1, 2)?
.contiguous()?;
k = k
.reshape((b_sz, q_len, self.num_kv_heads, self.head_dim))?
.transpose(1, 2)?
.contiguous()?;
} else if q.rank() == 3 {
q = q
.reshape((b_sz, self.num_heads, q_len, self.head_dim))?
.contiguous()?;
k = k
.reshape((b_sz, self.num_kv_heads, q_len, self.head_dim))?
.contiguous()?;
}
let mut attn_output = match &self.paged_attn {
Some(paged_attn) => match metadata {
Some(((key_cache, value_cache), input_metadata)) => paged_attn.forward(
&q,
&k,
&v,
attention_mask,
Some(key_cache),
Some(value_cache),
input_metadata,
None,
)?,
None => {
let mut input_metadata = PagedAttentionInputMetadata::dummy(q.device())?;
assert!(attention_mask.is_some());
paged_attn.forward(
&q,
&k,
&v,
attention_mask,
None,
None,
&mut input_metadata,
None,
)?
}
},
None => {
let (k, v, attn_mask) =
kv_cache.append_sliding_window(&k, &v, attention_mask, self.sliding_window)?;
Sdpa.run_attention(
&q,
&k,
&v,
attn_mask.as_ref(),
Some(flash_params),
&self.sdpa_params,
)?
}
};
if let Some(t) = self.q_proj.quantized_act_type() {
attn_output = attn_output.to_dtype(t)?;
}
attn_output = if attention_mask.is_some() {
attn_output.transpose(1, 2)?.reshape((b_sz, q_len, ()))?
} else {
attn_output.reshape((b_sz, q_len, ()))?
};
let mut res = MatMul.qmethod_matmul(&attn_output, &*self.o_proj)?;
if self.q_proj.quantized_act_type().is_some() {
res = res.to_dtype(original_dtype)?;
}
Ok(res)
}
}
#[derive(Clone)]
struct BlockSparseTop2MLP {
w1: Arc<dyn QuantMethod>,
w2: Arc<dyn QuantMethod>,
w3: Arc<dyn QuantMethod>,
act_fn: Activation,
}
impl BlockSparseTop2MLP {
fn new(cfg: &Config, vb: VarBuilder) -> Result<Self> {
let hidden_sz = cfg.hidden_size;
let intermediate_sz = cfg.intermediate_size;
let w1 = mistralrs_quant::linear_no_bias(
hidden_sz,
intermediate_sz,
&cfg.quantization_config,
vb.pp("w1"),
)?;
let w2 = mistralrs_quant::linear_no_bias(
intermediate_sz,
hidden_sz,
&cfg.quantization_config,
vb.pp("w2"),
)?;
let w3 = mistralrs_quant::linear_no_bias(
hidden_sz,
intermediate_sz,
&cfg.quantization_config,
vb.pp("w3"),
)?;
Ok(Self {
w1,
w2,
w3,
act_fn: cfg.hidden_act,
})
}
}
impl Module for BlockSparseTop2MLP {
fn forward(&self, xs: &Tensor) -> Result<Tensor> {
let original_dtype = xs.dtype();
let mut xs = xs.clone();
if let Some(t) = self.w1.quantized_act_type() {
xs = xs.to_dtype(t)?;
}
let lhs = MatMul.qmethod_matmul(&xs, &*self.w1)?.apply(&self.act_fn)?;
let rhs = MatMul.qmethod_matmul(&xs, &*self.w3)?;
let mut res = MatMul.qmethod_matmul(&(lhs * rhs)?, &*self.w2)?;
if self.w1.quantized_act_type().is_some() {
res = res.to_dtype(original_dtype)?;
}
Ok(res)
}
}
#[derive(Clone)]
struct SparseMoeBlock {
gate: Arc<dyn QuantMethod>,
experts: Vec<BlockSparseTop2MLP>,
num_experts_per_tok: usize,
}
impl SparseMoeBlock {
fn new(cfg: &Config, vb: VarBuilder) -> Result<Self> {
let gate = mistralrs_quant::linear_no_bias(
cfg.hidden_size,
cfg.num_local_experts,
&cfg.quantization_config,
vb.pp("gate"),
)?;
let mut experts = Vec::with_capacity(cfg.num_local_experts);
let vb = vb.pp("experts");
for idx in 0..cfg.num_local_experts {
let expert = BlockSparseTop2MLP::new(cfg, vb.pp(idx))?;
experts.push(expert)
}
Ok(SparseMoeBlock {
gate,
experts,
num_experts_per_tok: cfg.num_experts_per_tok,
})
}
}
impl Module for SparseMoeBlock {
fn forward(&self, xs: &Tensor) -> Result<Tensor> {
let (b_size, seq_len, hidden_dim) = xs.dims3()?;
let xs = xs.reshape(((), hidden_dim))?;
let original_dtype = xs.dtype();
let mut xs = xs.clone();
if let Some(t) = self.gate.quantized_act_type() {
xs = xs.to_dtype(t)?;
}
let mut router_logits = MatMul.qmethod_matmul(&xs, &*self.gate)?;
if self.gate.quantized_act_type().is_some() {
router_logits = router_logits.to_dtype(original_dtype)?;
}
let routing_weights = candle_nn::ops::softmax_last_dim(&router_logits)?;
let routing_weights = routing_weights.to_dtype(DType::F32)?.to_vec2::<f32>()?;
let mut top_x = vec![vec![]; self.experts.len()];
let mut selected_rws = vec![vec![]; self.experts.len()];
for (row_idx, rw) in routing_weights.iter().enumerate() {
let mut dst = (0..rw.len() as u32).collect::<Vec<u32>>();
dst.sort_by(|&i, &j| rw[j as usize].total_cmp(&rw[i as usize]));
let mut sum_routing_weights = 0f32;
for &expert_idx in dst.iter().take(self.num_experts_per_tok) {
let expert_idx = expert_idx as usize;
let routing_weight = rw[expert_idx];
sum_routing_weights += routing_weight;
top_x[expert_idx].push(row_idx as u32);
}
for &expert_idx in dst.iter().take(self.num_experts_per_tok) {
let expert_idx = expert_idx as usize;
let routing_weight = rw[expert_idx];
selected_rws[expert_idx].push(routing_weight / sum_routing_weights)
}
}
let mut ys = xs.zeros_like()?;
for (expert_idx, expert_layer) in self.experts.iter().enumerate() {
let top_x = &top_x[expert_idx];
if top_x.is_empty() {
continue;
}
let top_x = Tensor::new(top_x.as_slice(), xs.device())?;
let selected_rws =
Tensor::new(selected_rws[expert_idx].as_slice(), xs.device())?.reshape(((), 1))?;
let current_state = xs.index_select(&top_x, 0)?.reshape(((), hidden_dim))?;
let current_hidden_states = expert_layer.forward(¤t_state)?;
let current_hidden_states = current_hidden_states.broadcast_mul(&selected_rws)?;
ys = ys.index_add(&top_x, ¤t_hidden_states, 0)?;
}
let ys = ys.reshape((b_size, seq_len, hidden_dim))?;
Ok(ys)
}
}
struct DecoderLayer {
self_attn: Attention,
block_sparse_moe: SparseMoeBlock,
input_layernorm: RmsNorm,
post_attention_layernorm: RmsNorm,
}
impl DecoderLayer {
fn new(
rotary_emb: Arc<RotaryEmbedding>,
cfg: &Config,
vb: VarBuilder,
mapper: &dyn DeviceMapper,
layer_idx: usize,
loading_isq: bool,
paged_attn: Option<PagedAttention>,
) -> Result<Self> {
let self_attn = Attention::new(
rotary_emb,
cfg,
mapper.set_device(layer_idx, vb.pp("self_attn"), loading_isq),
paged_attn,
)?;
let block_sparse_moe = SparseMoeBlock::new(
cfg,
mapper.set_device(layer_idx, vb.pp("block_sparse_moe"), loading_isq),
)?;
let input_layernorm = RmsNorm::new(
cfg.hidden_size,
cfg.rms_norm_eps,
mapper.set_device(layer_idx, vb.pp("input_layernorm"), false),
)?;
let post_attention_layernorm = RmsNorm::new(
cfg.hidden_size,
cfg.rms_norm_eps,
mapper.set_device(layer_idx, vb.pp("post_attention_layernorm"), false),
)?;
Ok(Self {
self_attn,
block_sparse_moe,
input_layernorm,
post_attention_layernorm,
})
}
#[allow(clippy::too_many_arguments)]
fn forward(
&self,
xs: &Tensor,
attention_mask: Option<&Tensor>,
seqlen_offsets: &[usize],
start_offsets_kernel: Tensor,
kv_cache: &mut KvCache,
metadata: Option<((Tensor, Tensor), &mut PagedAttentionInputMetadata)>,
flash_params: &FlashParams,
) -> Result<Tensor> {
let residual = xs;
let xs = self.input_layernorm.forward(xs)?;
let xs = self.self_attn.forward(
&xs,
attention_mask,
seqlen_offsets,
start_offsets_kernel,
kv_cache,
metadata,
flash_params,
)?;
let xs = (xs + residual)?;
let residual = &xs;
let xs = xs
.apply(&self.post_attention_layernorm)?
.apply(&self.block_sparse_moe)?
.to_dtype(residual.dtype())?;
residual + xs
}
}
pub struct Model {
embed_tokens: candle_nn::Embedding,
layers: Vec<DecoderLayer>,
norm: RmsNorm,
lm_head: Arc<dyn QuantMethod>,
sliding_window: Option<usize>,
device: Device,
cache: EitherCache,
max_seq_len: usize,
mapper: Box<dyn DeviceMapper + Send + Sync>,
cfg: ModelConfigMetadata,
}
impl Model {
pub fn new(
cfg: &Config,
vb: VarBuilder,
is_gptx: bool,
normal_loading_metadata: NormalLoadingMetadata,
attention_mechanism: AttentionImplementation,
) -> Result<Self> {
if let Some(ref quant_cfg) = &cfg.quantization_config {
tracing::info!(
"Using {} quantization: {}.",
quant_cfg.quant_method.to_string(),
quant_cfg.get_bits_name(&vb)
);
}
let mapper = normal_loading_metadata.mapper;
let vb_m = vb.pp("model");
let embed_tokens = candle_nn::embedding(
cfg.vocab_size,
cfg.hidden_size,
mapper.set_nm_device(vb_m.pp("embed_tokens"), false),
)?;
let head_dim = cfg.hidden_size / cfg.num_attention_heads;
let mut ropes = HashMap::new();
for layer_idx in 0..cfg.num_hidden_layers {
let device = mapper
.device_for(layer_idx, false)
.unwrap_or(&normal_loading_metadata.real_device);
ropes.insert(
device.location(),
Arc::new(RotaryEmbedding::new(
cfg.rope_theta as f32,
head_dim,
cfg.max_position_embeddings,
device,
is_gptx,
vb_m.dtype(),
)?),
);
}
let mut layers = Vec::with_capacity(cfg.num_hidden_layers);
let vb_l = vb_m.pp("layers");
for layer_idx in
NiceProgressBar::<_, 'b'>(0..cfg.num_hidden_layers, "Loading repeating layers")
{
let device = mapper
.device_for(layer_idx, false)
.unwrap_or(&normal_loading_metadata.real_device);
let rotary_emb = ropes
.get(&device.location())
.expect("No RoPE for device location!")
.clone();
let paged_attn = match &attention_mechanism {
AttentionImplementation::Eager => None,
AttentionImplementation::PagedAttention => Some(PagedAttention::new(
cfg.num_attention_heads,
head_dim,
(1.0 / (head_dim as f64).sqrt()) as f32,
Some(cfg.num_key_value_heads),
cfg.sliding_window,
device,
None,
)?),
};
let layer = DecoderLayer::new(
rotary_emb.clone(),
cfg,
vb_l.pp(layer_idx),
&*mapper,
layer_idx,
normal_loading_metadata.loading_isq,
paged_attn,
)?;
layers.push(layer)
}
let norm = RmsNorm::new(
cfg.hidden_size,
cfg.rms_norm_eps,
mapper.set_nm_device(vb_m.pp("norm"), false),
)?;
let lm_head = if !cfg.tie_word_embeddings {
mistralrs_quant::linear_no_bias(
cfg.hidden_size,
cfg.vocab_size,
&None,
mapper.set_nm_device(vb.pp("lm_head"), normal_loading_metadata.loading_isq),
)?
} else {
Arc::new(UnquantLinear::new(QuantMethodConfig::Unquantized(
candle_nn::Linear::new(
mapper.cast_nm_device(
embed_tokens.embeddings(),
normal_loading_metadata.loading_isq,
)?,
None,
),
))?)
};
Ok(Self {
embed_tokens,
layers,
norm,
lm_head,
sliding_window: cfg.sliding_window,
device: normal_loading_metadata.real_device,
cache: EitherCache::Normal(NormalCache::new(
cfg.num_hidden_layers,
cfg.max_position_embeddings,
)),
max_seq_len: cfg.max_position_embeddings,
mapper,
cfg: ModelConfigMetadata {
num_layers: cfg.num_hidden_layers,
hidden_size: cfg.hidden_size,
num_kv_heads: cfg.num_key_value_heads,
num_attn_heads: cfg.num_attention_heads,
sliding_window: cfg.sliding_window,
k_head_dim: None,
v_head_dim: None,
},
})
}
pub fn forward(
&self,
input_ids: &Tensor,
seqlen_offsets: &[usize],
start_offsets_kernel: Tensor,
context_lens: Vec<(usize, usize)>,
mut metadata: Option<(Vec<(Tensor, Tensor)>, &mut PagedAttentionInputMetadata)>,
flash_params: &FlashParams,
) -> Result<Tensor> {
let mut xs = self.embed_tokens.forward(input_ids)?;
let cache = &mut self.cache.normal().0;
let attention_mask = CausalMasker.make_sliding_window_causal_mask_matrix(
input_ids,
metadata
.as_ref()
.map(|(_, _)| &seqlen_offsets as &dyn PastKvLenCache)
.unwrap_or(cache as &dyn PastKvLenCache),
self.sliding_window,
xs.dtype(),
self.cfg.num_attn_heads,
)?;
for (i, layer) in self.layers.iter().enumerate() {
xs = self.mapper.map(xs, i)?;
xs = layer.forward(
&xs,
attention_mask
.as_ref()
.map(|m| m.to_device(xs.device()).unwrap())
.as_ref(),
seqlen_offsets,
start_offsets_kernel.clone(),
&mut cache[i],
metadata
.as_mut()
.map(|(kv_cache, metadata)| (kv_cache[i].clone(), &mut **metadata)),
flash_params,
)?;
}
let xs = xs.to_device(&self.device)?;
let mut xs = xs.apply(&self.norm)?;
if let Some(t) = self.lm_head.quantized_act_type() {
xs = xs.to_dtype(t)?;
}
extract_logits(&MatMul.qmethod_matmul(&xs, &*self.lm_head)?, context_lens)
}
}
impl IsqModel for Model {
fn get_layers(
&mut self,
) -> (
Vec<(&mut Arc<dyn QuantMethod>, Option<usize>)>,
&dyn DeviceMapper,
) {
let mut tensors = Vec::new();
tensors.push((&mut self.lm_head, None));
for (i, layer) in self.layers.iter_mut().enumerate() {
tensors.push((&mut layer.self_attn.q_proj, Some(i)));
tensors.push((&mut layer.self_attn.k_proj, Some(i)));
tensors.push((&mut layer.self_attn.v_proj, Some(i)));
tensors.push((&mut layer.self_attn.o_proj, Some(i)));
tensors.push((&mut layer.block_sparse_moe.gate, Some(i)));
for expert in &mut layer.block_sparse_moe.experts {
tensors.push((&mut expert.w1, Some(i)));
tensors.push((&mut expert.w2, Some(i)));
tensors.push((&mut expert.w3, Some(i)));
}
}
(tensors, &*self.mapper)
}
fn residual_tensors(&self) -> Vec<(String, Tensor)> {
let uvb = UnVarBuilder::new();
let uvb_m = uvb.pp("model");
uvb_m.pp("embed_tokens").add(&self.embed_tokens);
uvb_m.pp("norm").add(&self.norm);
for (layer_idx, layer) in self.layers.iter().enumerate() {
let uvb_l = uvb_m.pp("layers").pp(layer_idx);
uvb_l.pp("input_layernorm").add(&layer.input_layernorm);
uvb_l
.pp("post_attention_layernorm")
.add(&layer.post_attention_layernorm);
}
uvb.to_safetensors()
}
}
impl NormalModel for Model {
fn forward(
&self,
input_ids: &Tensor,
seqlen_offsets: &[usize],
start_offsets_kernel: Tensor,
context_lens: Vec<(usize, usize)>,
_position_ids: Vec<usize>,
metadata: Option<(Vec<(Tensor, Tensor)>, &mut PagedAttentionInputMetadata)>,
flash_params: &FlashParams,
) -> Result<Tensor> {
self.forward(
input_ids,
seqlen_offsets,
start_offsets_kernel,
context_lens,
metadata,
flash_params,
)
}
fn xlora_forward(
&self,
_input_ids: &Tensor,
_input_ids_full: &Tensor,
_seqlen_offsets: &[usize],
_seqlen_offsets_full: &[usize],
_start_offsets_kernel: Tensor,
_start_offsets_kernel_full: Tensor,
_no_kv_cache: bool,
_non_granular_state: &Option<crate::xlora_models::NonGranularState>,
_context_lens: Vec<(usize, usize)>,
_position_ids: Vec<usize>,
_flash_params: &FlashParams,
_flash_params_full: &FlashParams,
) -> Result<Tensor> {
unimplemented!()
}
fn cache(&self) -> &EitherCache {
&self.cache
}
fn cache_mut(&mut self) -> &mut EitherCache {
&mut self.cache
}
fn device(&self) -> &Device {
&self.device
}
fn is_xlora(&self) -> bool {
false
}
fn max_seq_len(&self) -> usize {
self.max_seq_len
}
fn config(&self) -> &ModelConfigMetadata {
&self.cfg
}
}
impl AnyMoeBaseModelMixin for Model {}