Skip to content

Commit 872abf3

Browse files
2015arorasIsotr0py
authored andcommitted
[Model] Add Olmo3 model implementation (vllm-project#24534)
Signed-off-by: Shane A <shanea@allenai.org> Co-authored-by: Isotr0py <mozf@mail2.sysu.edu.cn>
1 parent d2f3920 commit 872abf3

File tree

7 files changed

+114
-14
lines changed

7 files changed

+114
-14
lines changed

docs/models/supported_models.md

Lines changed: 1 addition & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -389,6 +389,7 @@ th {
389389
| `NemotronHForCausalLM` | Nemotron-H | `nvidia/Nemotron-H-8B-Base-8K`, `nvidia/Nemotron-H-47B-Base-8K`, `nvidia/Nemotron-H-56B-Base-8K`, etc. | ✅︎ | ✅︎ | ✅︎ |
390390
| `OLMoForCausalLM` | OLMo | `allenai/OLMo-1B-hf`, `allenai/OLMo-7B-hf`, etc. | ✅︎ | ✅︎ | ✅︎ |
391391
| `OLMo2ForCausalLM` | OLMo2 | `allenai/OLMo-2-0425-1B`, etc. | ✅︎ | ✅︎ | ✅︎ |
392+
| `OLMo3ForCausalLM` | OLMo3 | TBA | ✅︎ | ✅︎ | ✅︎ |
392393
| `OLMoEForCausalLM` | OLMoE | `allenai/OLMoE-1B-7B-0924`, `allenai/OLMoE-1B-7B-0924-Instruct`, etc. | | ✅︎ | ✅︎ |
393394
| `OPTForCausalLM` | OPT, OPT-IML | `facebook/opt-66b`, `facebook/opt-iml-max-30b`, etc. | | ✅︎ | ✅︎ |
394395
| `OrionForCausalLM` | Orion | `OrionStarAI/Orion-14B-Base`, `OrionStarAI/Orion-14B-Chat`, etc. | | ✅︎ | ✅︎ |

tests/models/registry.py

Lines changed: 1 addition & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -301,6 +301,7 @@ def check_available_online(
301301
trust_remote_code=True),
302302
"OlmoForCausalLM": _HfExamplesInfo("allenai/OLMo-1B-hf"),
303303
"Olmo2ForCausalLM": _HfExamplesInfo("allenai/OLMo-2-0425-1B"),
304+
"Olmo3ForCausalLM": _HfExamplesInfo("shanearora/2025-sep-a-base-model"),
304305
"OlmoeForCausalLM": _HfExamplesInfo("allenai/OLMoE-1B-7B-0924-Instruct"),
305306
"OPTForCausalLM": _HfExamplesInfo("facebook/opt-125m",
306307
{"1b": "facebook/opt-iml-max-1.3b"}),

vllm/model_executor/models/olmo2.py

Lines changed: 28 additions & 14 deletions
Original file line numberDiff line numberDiff line change
@@ -52,10 +52,11 @@
5252
from vllm.model_executor.model_loader.weight_utils import default_weight_loader
5353
from vllm.model_executor.models.interfaces import SupportsLoRA, SupportsPP
5454
from vllm.model_executor.models.utils import (
55-
AutoWeightsLoader, is_pp_missing_parameter,
55+
AutoWeightsLoader, extract_layer_index, is_pp_missing_parameter,
5656
make_empty_intermediate_tensors_factory, make_layers, maybe_prefix)
5757
from vllm.model_executor.sampling_metadata import SamplingMetadata
5858
from vllm.sequence import IntermediateTensors
59+
from vllm.transformers_utils.configs import Olmo3Config
5960

6061

6162
class Olmo2Attention(nn.Module):
@@ -68,7 +69,7 @@ class Olmo2Attention(nn.Module):
6869
def __init__(self, *, vllm_config: VllmConfig, prefix: str = ""):
6970
super().__init__()
7071
self.config = vllm_config.model_config.hf_config
71-
assert isinstance(self.config, Olmo2Config)
72+
assert isinstance(self.config, (Olmo2Config, Olmo3Config))
7273

7374
hidden_size = self.config.hidden_size
7475
self.tp_size = get_tensor_model_parallel_world_size()
@@ -111,22 +112,35 @@ def __init__(self, *, vllm_config: VllmConfig, prefix: str = ""):
111112
self.q_norm = RMSNorm(self.config.hidden_size,
112113
eps=self.config.rms_norm_eps)
113114

114-
# Rotary embeddings.
115-
self.rotary_emb = get_rope(
116-
self.head_dim,
117-
rotary_dim=self.head_dim,
118-
max_position=self.max_position_embeddings,
119-
base=self.rope_theta, # type: ignore
120-
)
121115
self.scaling = self.head_dim**-0.5
116+
117+
layer_idx = extract_layer_index(prefix)
118+
sliding_window = None
119+
if ((layer_types := getattr(self.config, "layer_types", None))
120+
is not None and layer_types[layer_idx] == "sliding_attention"):
121+
sliding_window = self.config.sliding_window
122+
122123
self.attn = Attention(
123124
self.num_heads,
124125
self.head_dim,
125126
self.scaling,
126127
num_kv_heads=self.num_kv_heads,
127128
cache_config=vllm_config.cache_config,
128129
quant_config=vllm_config.quant_config,
129-
prefix=prefix,
130+
per_layer_sliding_window=sliding_window,
131+
prefix=f"{prefix}.attn",
132+
)
133+
134+
# Rotary embeddings. Rope scaling is only applied on full attention
135+
# layers.
136+
self.rope_scaling = (self.config.rope_scaling
137+
if sliding_window is None else None)
138+
self.rotary_emb = get_rope(
139+
self.head_dim,
140+
rotary_dim=self.head_dim,
141+
max_position=self.max_position_embeddings,
142+
base=self.rope_theta, # type: ignore
143+
rope_scaling=self.rope_scaling,
130144
)
131145

132146
# Attention output projection.
@@ -176,7 +190,7 @@ class Olmo2MLP(nn.Module):
176190
def __init__(self, *, vllm_config: VllmConfig, prefix: str = ""):
177191
super().__init__()
178192
config = vllm_config.model_config.hf_config
179-
assert isinstance(config, Olmo2Config)
193+
assert isinstance(config, (Olmo2Config, Olmo3Config))
180194
hidden_size = config.hidden_size
181195
intermediate_size = config.intermediate_size
182196

@@ -221,7 +235,7 @@ class Olmo2DecoderLayer(nn.Module):
221235
def __init__(self, *, vllm_config: VllmConfig, prefix: str = ""):
222236
super().__init__()
223237
config = vllm_config.model_config.hf_config
224-
assert isinstance(config, Olmo2Config)
238+
assert isinstance(config, (Olmo2Config, Olmo3Config))
225239
# Attention block.
226240
self.self_attn = Olmo2Attention(vllm_config=vllm_config,
227241
prefix=f"{prefix}.self_attn")
@@ -261,7 +275,7 @@ class Olmo2Model(nn.Module):
261275
def __init__(self, *, vllm_config: VllmConfig, prefix: str = ""):
262276
super().__init__()
263277
self.config = vllm_config.model_config.hf_config
264-
assert isinstance(self.config, Olmo2Config)
278+
assert isinstance(self.config, (Olmo2Config, Olmo3Config))
265279

266280
self.embed_tokens = VocabParallelEmbedding(
267281
self.config.vocab_size,
@@ -376,7 +390,7 @@ class Olmo2ForCausalLM(nn.Module, SupportsPP, SupportsLoRA):
376390
def __init__(self, *, vllm_config: VllmConfig, prefix: str = ""):
377391
super().__init__()
378392
config = vllm_config.model_config.hf_config
379-
assert isinstance(config, Olmo2Config)
393+
assert isinstance(config, (Olmo2Config, Olmo3Config))
380394
self.config = config
381395
self.model = Olmo2Model(vllm_config=vllm_config,
382396
prefix=maybe_prefix(prefix, "model"))

vllm/model_executor/models/registry.py

Lines changed: 1 addition & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -120,6 +120,7 @@
120120
"NemotronHForCausalLM": ("nemotron_h", "NemotronHForCausalLM"),
121121
"OlmoForCausalLM": ("olmo", "OlmoForCausalLM"),
122122
"Olmo2ForCausalLM": ("olmo2", "Olmo2ForCausalLM"),
123+
"Olmo3ForCausalLM": ("olmo2", "Olmo2ForCausalLM"),
123124
"OlmoeForCausalLM": ("olmoe", "OlmoeForCausalLM"),
124125
"OPTForCausalLM": ("opt", "OPTForCausalLM"),
125126
"OrionForCausalLM": ("orion", "OrionForCausalLM"),

vllm/transformers_utils/config.py

Lines changed: 1 addition & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -75,6 +75,7 @@ def __getitem__(self, key):
7575
eagle="EAGLEConfig",
7676
speculators="SpeculatorsConfig",
7777
nemotron="NemotronConfig",
78+
olmo3="Olmo3Config",
7879
ovis="OvisConfig",
7980
ultravox="UltravoxConfig",
8081
step3_vl="Step3VLConfig",

vllm/transformers_utils/configs/__init__.py

Lines changed: 2 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -23,6 +23,7 @@
2323
from vllm.transformers_utils.configs.nemotron import NemotronConfig
2424
from vllm.transformers_utils.configs.nemotron_h import NemotronHConfig
2525
from vllm.transformers_utils.configs.nemotron_vl import Nemotron_Nano_VL_Config
26+
from vllm.transformers_utils.configs.olmo3 import Olmo3Config
2627
from vllm.transformers_utils.configs.ovis import OvisConfig
2728
from vllm.transformers_utils.configs.qwen3_next import Qwen3NextConfig
2829
from vllm.transformers_utils.configs.speculators.base import SpeculatorsConfig
@@ -45,6 +46,7 @@
4546
"NemotronConfig",
4647
"NemotronHConfig",
4748
"Nemotron_Nano_VL_Config",
49+
"Olmo3Config",
4850
"OvisConfig",
4951
"SpeculatorsConfig",
5052
"UltravoxConfig",
Lines changed: 80 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,80 @@
1+
# SPDX-License-Identifier: Apache-2.0
2+
# SPDX-FileCopyrightText: Copyright contributors to the vLLM project
3+
4+
from transformers.configuration_utils import PretrainedConfig
5+
6+
7+
class Olmo3Config(PretrainedConfig):
8+
9+
model_type = "olmo3"
10+
keys_to_ignore_at_inference = ["past_key_values"]
11+
12+
def __init__(
13+
self,
14+
vocab_size=50304,
15+
hidden_size=4096,
16+
intermediate_size=11008,
17+
num_hidden_layers=32,
18+
num_attention_heads=32,
19+
num_key_value_heads=None,
20+
hidden_act="silu",
21+
max_position_embeddings=2048,
22+
initializer_range=0.02,
23+
use_cache=True,
24+
pad_token_id=1,
25+
bos_token_id=None,
26+
eos_token_id=50279,
27+
tie_word_embeddings=False,
28+
rope_theta=10000.0,
29+
rope_scaling=None,
30+
attention_bias=False,
31+
attention_dropout=0.0,
32+
rms_norm_eps=1e-5,
33+
sliding_window=4096,
34+
layer_types=None,
35+
**kwargs,
36+
):
37+
# This model uses Olmo3ForCausalLM in transformers but Olmo2ForCausalLM
38+
# in vLLM.
39+
if "architectures" not in kwargs:
40+
kwargs["architectures"] = ["Olmo2ForCausalLM"]
41+
elif "Olmo3ForCausalLM" in kwargs["architectures"]:
42+
kwargs["architectures"].remove("Olmo3ForCausalLM")
43+
kwargs["architectures"].append("Olmo2ForCausalLM")
44+
45+
super().__init__(
46+
pad_token_id=pad_token_id,
47+
bos_token_id=bos_token_id,
48+
eos_token_id=eos_token_id,
49+
tie_word_embeddings=tie_word_embeddings,
50+
**kwargs,
51+
)
52+
self.vocab_size = vocab_size
53+
self.max_position_embeddings = max_position_embeddings
54+
self.hidden_size = hidden_size
55+
self.intermediate_size = intermediate_size
56+
self.num_hidden_layers = num_hidden_layers
57+
self.num_attention_heads = num_attention_heads
58+
59+
# for backward compatibility
60+
if num_key_value_heads is None:
61+
num_key_value_heads = num_attention_heads
62+
63+
self.num_key_value_heads = num_key_value_heads
64+
self.hidden_act = hidden_act
65+
self.initializer_range = initializer_range
66+
self.use_cache = use_cache
67+
self.rope_theta = rope_theta
68+
self.rope_scaling = rope_scaling
69+
self.attention_bias = attention_bias
70+
self.attention_dropout = attention_dropout
71+
72+
self.rms_norm_eps = rms_norm_eps
73+
74+
self.sliding_window = sliding_window
75+
self.layer_types = layer_types
76+
if self.layer_types is None:
77+
self.layer_types = [
78+
"sliding_attention" if (i + 1) % 4 != 0 else "full_attention"
79+
for i in range(self.num_hidden_layers)
80+
]

0 commit comments

Comments
 (0)