test_initialization.py 4.74 KB
Newer Older
1
# SPDX-License-Identifier: Apache-2.0
2
# SPDX-FileCopyrightText: Copyright contributors to the vLLM project
3

4
5
6
7
8
9
from unittest.mock import patch

import pytest
from transformers import PretrainedConfig

from vllm import LLM
10
from vllm.engine.llm_engine import LLMEngine as V0LLMEngine
11
12
from vllm.utils import GiB_bytes
from vllm.v1.core.kv_cache_utils import get_kv_cache_config
13
from vllm.v1.engine.core import EngineCore as V1EngineCore
14

15
from ..utils import create_new_process_for_each_test
16
17
18
19
from .registry import HF_EXAMPLE_MODELS


@pytest.mark.parametrize("model_arch", HF_EXAMPLE_MODELS.get_supported_archs())
20
@create_new_process_for_each_test()
21
def test_can_initialize(model_arch: str, monkeypatch: pytest.MonkeyPatch):
22
23
24
25
26
27
28
    """The reason for using create_new_process_for_each_test is to avoid 
    the WARNING: 
        "We must use the 'spawn' multiprocessing start method. Overriding 
        VLLM_WORKER_MULTIPROC_METHOD to 'spawn'."
    The spawn process causes the _initialize_kv_caches_v1 function below to 
    become ineffective.
    """
29
    model_info = HF_EXAMPLE_MODELS.get_hf_info(model_arch)
30
31
    model_info.check_available_online(on_fail="skip")
    model_info.check_transformers_version(on_fail="skip")
32

33
    # FIXME: Possible memory leak in the previous tests?
34
35
    if model_arch in ("Glm4vForConditionalGeneration",
                      "GraniteSpeechForConditionalGeneration",
36
                      "KimiVLForConditionalGeneration"):
37
38
        pytest.skip("Avoid OOM")

39
    # Avoid OOM and reduce initialization time by only using 1 layer
40
    def hf_overrides(hf_config: PretrainedConfig) -> PretrainedConfig:
41
        hf_config.update(model_info.hf_overrides)
42

43
        text_config = hf_config.get_text_config()
44

45
46
        # Ensure at least 2 expert per group
        # Since `grouped_topk` assums top-2
47
48
        n_group = getattr(text_config, 'n_group', None)
        num_experts = n_group * 2 if n_group is not None else 2
49

50
51
52
53
54
        # we use three layers for Gemma-3n to check
        # both normal layer and kv_shared_layer
        num_hidden_layers = (3 if model_arch
                             == "Gemma3nForConditionalGeneration" else 1)

55
56
        text_config.update({
            "num_layers": 1,
57
            "num_hidden_layers": num_hidden_layers,
58
            "num_experts": num_experts,
59
            "num_experts_per_tok": 2,
60
61
62
63
64
            "num_local_experts": num_experts,
            # Otherwise there will not be any expert layers
            "first_k_dense_replace": 0,
            # To avoid OOM on DeepSeek-V3
            "n_routed_experts": num_experts,
65
66
            # For Gemma-3n
            "num_kv_shared_layers": 1,
67
68
        })

69
70
71
72
73
74
        if hasattr(hf_config, "vision_config"):
            hf_config.vision_config.update({
                "num_layers": 1,
                "num_hidden_layers": 1,
            })

75
76
77
78
79
80
81
        # e.g.: ibm-granite/granite-speech-3.3-2b
        if hasattr(hf_config, "encoder_config"):
            hf_config.encoder_config.update({
                "num_layers": 1,
                "num_hidden_layers": 1,
            })

82
83
84
        return hf_config

    # Avoid calling model.forward()
85
    def _initialize_kv_caches_v0(self) -> None:
86
87
88
        self.cache_config.num_gpu_blocks = 0
        self.cache_config.num_cpu_blocks = 0

89
90
91
92
93
    def _initialize_kv_caches_v1(self, vllm_config):
        kv_cache_specs = self.model_executor.get_kv_cache_specs()
        scheduler_kv_cache_config = get_kv_cache_config(
            vllm_config,
            kv_cache_specs[0],
94
            10 * GiB_bytes,
95
96
97
98
        )

        # gpu_blocks (> 0), cpu_blocks, scheduler_kv_cache_config
        return 1, 0, scheduler_kv_cache_config
99
100
101
102

    with (patch.object(V0LLMEngine, "_initialize_kv_caches",
                       _initialize_kv_caches_v0),
          patch.object(V1EngineCore, "_initialize_kv_caches",
103
104
105
                       _initialize_kv_caches_v1), monkeypatch.context() as m):
        if model_info.v0_only:
            m.setenv("VLLM_USE_V1", "0")
106
107
108
        if model_arch == "Phi4FlashForCausalLM":
            # Phi4FlashForCausalLM only supports DIFFERENTIAL_FLASH_ATTN backend
            m.setenv("VLLM_ATTENTION_BACKEND", "DIFFERENTIAL_FLASH_ATTN")
109
        LLM(
110
            model_info.default,
111
112
            tokenizer=model_info.tokenizer,
            tokenizer_mode=model_info.tokenizer_mode,
113
            revision=model_info.revision,
114
115
116
117
            speculative_config={
                "model": model_info.speculative_model,
                "num_speculative_tokens": 1,
            } if model_info.speculative_model else None,
118
            trust_remote_code=model_info.trust_remote_code,
119
            max_model_len=model_info.max_model_len,
120
121
            # these tests seem to produce leftover memory
            gpu_memory_utilization=0.80,
122
123
124
            load_format="dummy",
            hf_overrides=hf_overrides,
        )