utils.py 6.03 KB
Newer Older
1
2
3
4
5
6
# SPDX-License-Identifier: Apache-2.0
# SPDX-FileCopyrightText: Copyright contributors to the vLLM project
from typing import Optional, Union

import torch

7
8
9
10
11
12
13
14
15
16
17
18
19
from vllm.config import (
    CacheConfig,
    KVTransferConfig,
    ModelConfig,
    SchedulerConfig,
    SpeculativeConfig,
    VllmConfig,
)
from vllm.multimodal.inputs import (
    MultiModalFeatureSpec,
    MultiModalKwargsItem,
    PlaceholderRange,
)
20
from vllm.sampling_params import SamplingParams
21
from vllm.utils import sha256
22
from vllm.v1.core.kv_cache_utils import get_request_block_hasher, init_none_hash
23
24
from vllm.v1.core.sched.async_scheduler import AsyncScheduler
from vllm.v1.core.sched.scheduler import Scheduler
25
26
27
28
29
from vllm.v1.kv_cache_interface import (
    FullAttentionSpec,
    KVCacheConfig,
    KVCacheGroupSpec,
)
30
31
32
33
34
35
36
37
38
39
40
41
42
43
44
45
46
47
48
49
50
from vllm.v1.request import Request
from vllm.v1.structured_output import StructuredOutputManager

EOS_TOKEN_ID = 50256


def create_scheduler(
    model: str = "facebook/opt-125m",
    max_num_seqs: int = 16,
    max_num_batched_tokens: int = 8192,
    enable_prefix_caching: Optional[bool] = None,
    long_prefill_token_threshold: int = 0,
    disable_chunked_mm_input: bool = False,
    use_kv_connector: bool = False,
    num_blocks: int = 10000,
    block_size: int = 16,
    max_model_len: Optional[int] = None,
    num_speculative_tokens: Optional[int] = None,
    skip_tokenizer_init: bool = False,
    async_scheduling: bool = False,
) -> Union[Scheduler, AsyncScheduler]:
51
    """Create scheduler under test.
52
53
54
55
56
57
58
59
60
61
62

    Args:
      model: model under test
      max_num_seqs: max sequences to schedule
      max_num_batch_tokens: max num tokens to batch
      enable_prefix_caching: optionally force APC config
                             (True/False) or use default
                             (None)

    Returns:
      {class}`Scheduler` instance
63
    """
64
65
66
67
68
69
70
71
72
73
74
75
76
77
78
79
80
81
82
    if max_model_len is None:
        max_model_len = max_num_batched_tokens
    scheduler_config = SchedulerConfig(
        max_num_seqs=max_num_seqs,
        max_num_batched_tokens=max_num_batched_tokens,
        max_model_len=max_model_len,
        long_prefill_token_threshold=long_prefill_token_threshold,
        disable_chunked_mm_input=disable_chunked_mm_input,
        enable_chunked_prefill=True,
        async_scheduling=async_scheduling,
    )
    model_config = ModelConfig(
        model=model,
        trust_remote_code=True,
        dtype="float16",
        seed=42,
        skip_tokenizer_init=skip_tokenizer_init,
    )
    # Cache config, optionally force APC
83
84
85
86
87
    kwargs_cache = (
        {}
        if enable_prefix_caching is None
        else {"enable_prefix_caching": enable_prefix_caching}
    )
88
89
90
91
92
93
94
    cache_config = CacheConfig(
        block_size=block_size,
        gpu_memory_utilization=0.9,
        swap_space=0,
        cache_dtype="auto",
        **kwargs_cache,
    )
95
96
97
98
99
100
101
102
103
    kv_transfer_config = (
        KVTransferConfig(
            kv_connector="SharedStorageConnector",
            kv_role="kv_both",
            kv_connector_extra_config={"shared_storage_path": "local_storage"},
        )
        if use_kv_connector
        else None
    )
104
105
106
107

    speculative_config: Optional[SpeculativeConfig] = None
    if num_speculative_tokens is not None:
        speculative_config = SpeculativeConfig(
108
109
            model="ngram", num_speculative_tokens=num_speculative_tokens
        )
110
111
112
113
114
115
116
117
118
119
120
121

    vllm_config = VllmConfig(
        scheduler_config=scheduler_config,
        model_config=model_config,
        cache_config=cache_config,
        kv_transfer_config=kv_transfer_config,
        speculative_config=speculative_config,
    )
    kv_cache_config = KVCacheConfig(
        num_blocks=num_blocks,  # A large number of blocks to hold all requests
        kv_cache_tensors=[],
        kv_cache_groups=[
122
123
124
            KVCacheGroupSpec(
                ["layer"], FullAttentionSpec(block_size, 1, 1, torch.float32, False)
            )
125
126
127
128
129
130
131
        ],
    )
    cache_config.num_gpu_blocks = num_blocks
    scheduler_cls = AsyncScheduler if async_scheduling else Scheduler
    return scheduler_cls(
        vllm_config=vllm_config,
        kv_cache_config=kv_cache_config,
132
        block_size=block_size,
133
134
135
136
137
        log_stats=True,
        structured_output_manager=StructuredOutputManager(vllm_config),
    )


138
139
140
_none_hash_initialized = False


141
142
143
def create_requests(
    num_requests: int,
    num_tokens: int = 10,
144
    mm_positions: Optional[list[list[PlaceholderRange]]] = None,
145
146
147
148
    max_tokens: int = 16,
    stop_token_ids: Optional[list[int]] = None,
    prompt_logprobs: Optional[int] = None,
    same_prompt: bool = False,
149
    block_size: int = 16,
150
) -> list[Request]:
151
152
    global _none_hash_initialized
    if not _none_hash_initialized:
153
        init_none_hash(sha256)
154
155
        _none_hash_initialized = True

156
    block_hasher = get_request_block_hasher(block_size, sha256)
157
158
159
160
161
162
    sampling_params = SamplingParams(
        ignore_eos=False,
        max_tokens=max_tokens,
        stop_token_ids=stop_token_ids,
        prompt_logprobs=prompt_logprobs,
    )
163
164
    requests = []
    for i in range(num_requests):
165
        mm_features = []
166
167
        if mm_positions is not None:
            mm_position = mm_positions[i]
168
169
170
171
172
173
174
175
            for j, position in enumerate(mm_position):
                # Dummy hash for each mm item should be unique
                # since encoder cache tracks entries by hash
                identifier = f"hash{i}_{j}"
                mm_feature = MultiModalFeatureSpec(
                    data=MultiModalKwargsItem.dummy("dummy_m"),
                    mm_position=position,
                    identifier=identifier,
176
177
                    modality="image",
                )
178
179
                mm_features.append(mm_feature)

180
        prompt_token_ids = [0] * num_tokens if same_prompt else [i] * num_tokens
181
182
183
184
185
        request = Request(
            request_id=f"{i}",
            prompt_token_ids=prompt_token_ids,
            sampling_params=sampling_params,
            pooling_params=None,
186
            mm_features=mm_features if mm_features else None,
187
            eos_token_id=EOS_TOKEN_ID,
188
            block_hasher=block_hasher,
189
190
191
        )
        requests.append(request)
    return requests