"benchmarks/kernels/benchmark_moe_int4.py" did not exist on "3a434b07edc42a7466fb1ac536f3beb9470f9416"
registry.py 7.74 KB
Newer Older
1
# SPDX-License-Identifier: Apache-2.0
2
# SPDX-FileCopyrightText: Copyright contributors to the vLLM project
3
from collections.abc import Mapping
4
from dataclasses import dataclass
5
from typing import TYPE_CHECKING, Any, NamedTuple, Optional, Union
6

7
import torch
8
from packaging.version import Version
9
from transformers import BatchFeature, PretrainedConfig, ProcessorMixin
10
from transformers import __version__ as TRANSFORMERS_VERSION
11
from typing_extensions import TypeVar
12

13
14
from vllm.jsontree import JSONTree, json_map_leaves
from vllm.logger import init_logger
15
from vllm.transformers_utils.processor import cached_processor_from_config
16
from vllm.transformers_utils.tokenizer import AnyTokenizer
17
from vllm.utils import resolve_mm_processor_kwargs
18
19

if TYPE_CHECKING:
20
    from vllm.config import ModelConfig
21
22
    from vllm.multimodal import (MultiModalDataDict, MultiModalPlaceholderDict,
                                 MultiModalRegistry)
23
24
    from vllm.sequence import SequenceData

25
26
27
_T = TypeVar("_T")
_C = TypeVar("_C", bound=PretrainedConfig, default=PretrainedConfig)
_P = TypeVar("_P", bound=ProcessorMixin, default=ProcessorMixin)
28

29
30
logger = init_logger(__name__)

31

32
33
34
35
36
37
38
39
40
41
@dataclass(frozen=True)
class InputContext:
    """
    Contains information about the model which may be used to
    modify the inputs.
    """

    model_config: "ModelConfig"
    """The configuration of the model."""

42
43
    def get_hf_config(
        self,
44
        typ: Union[type[_C], tuple[type[_C], ...]] = PretrainedConfig,
45
        /,
46
    ) -> _C:
47
48
        """
        Get the HuggingFace configuration
49
        (`transformers.PretrainedConfig`) of the model,
50
51
52
        additionally checking its type.

        Raises:
53
            TypeError: If the configuration is not of the specified type.
54
55
        """
        hf_config = self.model_config.hf_config
56
        if not isinstance(hf_config, typ):
57
            raise TypeError("Invalid type of HuggingFace config. "
58
                            f"Expected type: {typ}, but "
59
60
61
62
                            f"found type: {type(hf_config)}")

        return hf_config

63
    def get_hf_image_processor_config(self) -> dict[str, Any]:
64
65
66
67
68
        """
        Get the HuggingFace image processor configuration of the model.
        """
        return self.model_config.hf_image_processor_config

69
70
71
72
73
74
75
76
77
78
79
80
81
    def get_mm_config(self):
        """
        Get the multimodal config of the model.

        Raises:
            RuntimeError: If the model is not a multimodal model.
        """
        mm_config = self.model_config.multimodal_config
        if mm_config is None:
            raise RuntimeError("Not a multimodal model")

        return mm_config

82
83
    def get_hf_processor(
        self,
84
        typ: Union[type[_P], tuple[type[_P], ...]] = ProcessorMixin,
85
86
        /,
        **kwargs: object,
87
    ) -> _P:
88
89
        """
        Get the HuggingFace processor
90
        (`transformers.ProcessorMixin`) of the model,
91
92
93
94
95
        additionally checking its type.

        Raises:
            TypeError: If the processor is not of the specified type.
        """
96
97
98
99
100
101
102
103
104
105
106
107
108
109
110
111
        return cached_processor_from_config(
            self.model_config,
            processor_cls=typ,
            **kwargs,
        )

    def init_processor(
        self,
        typ: type[_T],
        /,
        **kwargs: object,
    ) -> _T:
        """
        Initialize a HuggingFace-like processor class, merging the
        keyword arguments with those in the model's configuration.
        """
112
113
        mm_config = self.model_config.get_multimodal_config()
        base_kwargs = mm_config.mm_processor_kwargs
114
115
116
117
118
        if base_kwargs is None:
            base_kwargs = {}

        merged_kwargs = {**base_kwargs, **kwargs}

119
        return typ(**merged_kwargs)
120

121

122
123
124
125
126
@dataclass(frozen=True)
class InputProcessingContext(InputContext):
    tokenizer: AnyTokenizer
    """The tokenizer used to tokenize the inputs."""

127
128
    def get_hf_processor(
        self,
129
        typ: Union[type[_P], tuple[type[_P], ...]] = ProcessorMixin,
130
131
        /,
        **kwargs: object,
132
    ) -> _P:
133
134
135
136
137
        # Transformers 4.53.0 has issue with passing tokenizer to
        # initialize processor. We disable it for this version.
        # See: https://github.com/vllm-project/vllm/issues/20224
        if Version(TRANSFORMERS_VERSION) != Version("4.53.0"):
            kwargs["tokenizer"] = self.tokenizer
138
139
140
        return super().get_hf_processor(
            typ,
            **kwargs,
141
142
        )

143
    def call_hf_processor(
144
145
        self,
        hf_processor: ProcessorMixin,
146
147
        data: Mapping[str, object],
        kwargs: Mapping[str, object] = {},
148
    ) -> Union[BatchFeature, JSONTree]:
149
        """
150
151
        Call `hf_processor` on the prompt `data`
        (text, image, audio...) with configurable options `kwargs`.
152
        """
153
154
        assert callable(hf_processor)

155
156
        mm_config = self.model_config.get_multimodal_config()
        base_kwargs = mm_config.mm_processor_kwargs
157
158
159
        if base_kwargs is None:
            base_kwargs = {}

160
        merged_kwargs = resolve_mm_processor_kwargs(
161
            base_kwargs,
162
            kwargs,
163
            hf_processor,
164
165
            requires_kw_only=False,
            allow_var_kwargs=True,
166
        )
167

168
169
170
171
172
173
        def maybe_cast_dtype(x):
            # This mimics the behavior of transformers.BatchFeature
            if isinstance(x, torch.Tensor) and x.is_floating_point():
                return x.to(dtype=self.model_config.dtype)
            return x

174
        try:
175
176
177
            output = hf_processor(**data, **merged_kwargs, return_tensors="pt")
            # this emulates output.to(dtype=self.model_config.dtype)
            if isinstance(output, BatchFeature):
178
                cast_output = json_map_leaves(maybe_cast_dtype, output.data)
179
180
                return BatchFeature(cast_output)

181
182
            cast_output = json_map_leaves(maybe_cast_dtype, output)

183
184
185
186
187
188
            logger.warning_once(
                f"{type(hf_processor).__name__} did not return `BatchFeature`. "
                "Make sure to match the behaviour of `ProcessorMixin` when "
                "implementing custom processors.")
            return cast_output

189
190
191
192
        except Exception as exc:
            msg = (f"Failed to apply {type(hf_processor).__name__} "
                   f"on data={data} with kwargs={merged_kwargs}")

193
            raise ValueError(msg) from exc
194

195

196
class DummyData(NamedTuple):
197
198
199
200
201
    """
    Dummy data used for profiling.

    Note: This is only used in V0.
    """
202
203
204
205
206
207

    seq_data: "SequenceData"
    multi_modal_data: Optional["MultiModalDataDict"] = None
    multi_modal_placeholders: Optional["MultiModalPlaceholderDict"] = None


208
209
class InputRegistry:
    """
210
    Note: This is only used in V0.
211
212
    """

213
214
215
216
217
    def dummy_data_for_profiling(
        self,
        model_config: "ModelConfig",
        seq_len: int,
        mm_registry: "MultiModalRegistry",
218
        is_encoder_data: bool = False,
219
    ) -> DummyData:
220
221
222
223
224
225
        """
        Create dummy data for profiling the memory usage of a model.

        The model is identified by ``model_config``.
        """
        # Avoid circular import
226
        from vllm.sequence import SequenceData
227

228
229
230
        if not model_config.is_multimodal_model:
            seq_data = SequenceData.from_prompt_token_counts((0, seq_len))
            return DummyData(seq_data=seq_data)
231

232
233
234
235
236
237
        # Encoder dummy data does not contain multi-modal data
        if is_encoder_data:
            enc_data = mm_registry.get_encoder_dummy_data(
                model_config, seq_len)
            seq_data = SequenceData.from_seqs(enc_data.prompt_token_ids)
            return DummyData(seq_data=seq_data)
238

239
        dec_data = mm_registry.get_decoder_dummy_data(model_config, seq_len)
240

241
242
243
244
        return DummyData(
            seq_data=SequenceData.from_seqs(dec_data.prompt_token_ids),
            multi_modal_data=dec_data.multi_modal_data,
            multi_modal_placeholders=dec_data.multi_modal_placeholders,
245
        )