registry.py 14.6 KB
Newer Older
1
import functools
2
from collections import UserDict
3
4
5
from dataclasses import dataclass
from typing import (TYPE_CHECKING, Any, Dict, Generic, Mapping, Optional,
                    Protocol, Sequence, Type, TypeVar)
6

7
8
9
import torch.nn as nn

from vllm.inputs import InputProcessingContext
10
from vllm.logger import init_logger
11
from vllm.transformers_utils.tokenizer import AnyTokenizer
12
from vllm.utils import ClassRegistry
13

14
from .audio import AudioPlugin
15
from .base import MultiModalInputMapper, MultiModalPlugin, MultiModalTokensCalc
16
from .image import ImagePlugin
17
from .inputs import MultiModalDataDict, MultiModalKwargs, NestedTensors
18
19
20
from .processing import (BaseMultiModalProcessor, BaseProcessingInfo,
                         ProcessingCache)
from .profiling import BaseDummyInputsBuilder
21
from .utils import cached_get_tokenizer
22
from .video import VideoPlugin
23

24
25
26
if TYPE_CHECKING:
    from vllm.config import ModelConfig

27
28
logger = init_logger(__name__)

29
30
31
# TODO: Tune the MM cache size
MM_CACHE_SIZE = 256

32
N = TypeVar("N", bound=Type[nn.Module])
33
34
_I = TypeVar("_I", bound=BaseProcessingInfo)
_I_co = TypeVar("_I_co", bound=BaseProcessingInfo, covariant=True)
35
36


37
class ProcessingInfoFactory(Protocol[_I_co]):
38
39
40
41
42
    """Constructs a :class:`MultiModalProcessor` instance from the context."""

    def __call__(
        self,
        ctx: InputProcessingContext,
43
44
45
46
47
48
49
50
51
52
53
54
55
56
57
58
59
60
61
62
    ) -> _I_co:
        ...


class DummyInputsBuilderFactory(Protocol[_I]):
    """
    Constructs a :class:`BaseDummyInputsBuilder` instance from the context.
    """

    def __call__(self, info: _I) -> BaseDummyInputsBuilder[_I]:
        ...


class MultiModalProcessorFactory(Protocol[_I]):
    """Constructs a :class:`MultiModalProcessor` instance from the context."""

    def __call__(
        self,
        info: _I,
        dummy_inputs: BaseDummyInputsBuilder[_I],
63
64
        *,
        cache: Optional[ProcessingCache] = None,
65
    ) -> BaseMultiModalProcessor[_I]:
66
        ...
67

68

69
70
71
72
73
74
75
76
77
78
79
80
81
82
83
84
85
@dataclass(frozen=True)
class _ProcessorFactories(Generic[_I]):
    info: ProcessingInfoFactory[_I]
    processor: MultiModalProcessorFactory[_I]
    dummy_inputs: DummyInputsBuilderFactory[_I]

    def build_processor(
        self,
        ctx: InputProcessingContext,
        *,
        cache: Optional[ProcessingCache] = None,
    ):
        info = self.info(ctx)
        dummy_inputs_builder = self.dummy_inputs(info)
        return self.processor(info, dummy_inputs_builder, cache=cache)


86
class _MultiModalLimits(UserDict["ModelConfig", Dict[str, int]]):
87
88
89
90
91
    """
    Wraps `_limits_by_model` for a more informative error message
    when attempting to access a model that does not exist.
    """

92
    def __getitem__(self, key: "ModelConfig") -> Dict[str, int]:
93
94
95
96
97
98
99
100
        try:
            return super().__getitem__(key)
        except KeyError as exc:
            msg = (f"Cannot find `mm_limits` for model={key.model}. Did you "
                   "forget to call `init_mm_limits_per_prompt`?")
            raise KeyError(msg) from exc


101
102
class MultiModalRegistry:
    """
103
104
    A registry that dispatches data processing to the
    :class:`~vllm.multimodal.MultiModalPlugin` for each modality.
105
106
    """

107
    DEFAULT_PLUGINS = (ImagePlugin(), AudioPlugin(), VideoPlugin())
108

109
    def __init__(
110
111
112
113
            self,
            *,
            plugins: Sequence[MultiModalPlugin] = DEFAULT_PLUGINS) -> None:
        self._plugins = {p.get_data_key(): p for p in plugins}
114

115
        self._processor_factories = ClassRegistry[nn.Module,
116
                                                  _ProcessorFactories]()
117

118
119
120
121
122
        # This is used for non-multimodal models
        self._disabled_limits_per_plugin = {k: 0 for k in self._plugins}

        self._limits_by_model = _MultiModalLimits()

123
124
        self._processing_cache = ProcessingCache(MM_CACHE_SIZE)

125
    def register_plugin(self, plugin: MultiModalPlugin) -> None:
126
127
128
        """
        Register a multi-modal plugin so it can be recognized by vLLM.
        """
129
        data_type_key = plugin.get_data_key()
130

131
        if data_type_key in self._plugins:
132
133
            logger.warning(
                "A plugin is already registered for data type %s, "
134
                "and will be overwritten by the new plugin %s.", data_type_key,
135
136
                plugin)

137
        self._plugins[data_type_key] = plugin
138

139
140
141
142
    def _get_plugin(self, data_type_key: str):
        plugin = self._plugins.get(data_type_key)
        if plugin is not None:
            return plugin
143

144
        msg = f"Unknown multi-modal data type: {data_type_key}"
145
146
        raise NotImplementedError(msg)

147
    def register_input_mapper(
148
        self,
149
        data_type_key: str,
150
        mapper: Optional[MultiModalInputMapper] = None,
151
    ):
152
        """
153
        Register an input mapper for a specific modality to a model class.
154

155
        See :meth:`MultiModalPlugin.register_input_mapper` for more details.
156
        """
157
        return self._get_plugin(data_type_key).register_input_mapper(mapper)
158

159
    def register_image_input_mapper(
160
        self,
161
        mapper: Optional[MultiModalInputMapper] = None,
162
    ):
163
        """
164
        Register an input mapper for image data to a model class.
165

166
        See :meth:`MultiModalPlugin.register_input_mapper` for more details.
167
        """
168
        return self.register_input_mapper("image", mapper)
169

170
171
    def map_input(
        self,
172
        model_config: "ModelConfig",
173
174
        data: MultiModalDataDict,
        mm_processor_kwargs: Optional[Dict[str, Any]] = None,
175
    ) -> MultiModalKwargs:
176
        """
177
        Apply an input mapper to the data passed to the model.
178
179
180
181
182

        The data belonging to each modality is passed to the corresponding
        plugin which in turn converts the data into into keyword arguments
        via the input mapper registered for that model.

183
        See :meth:`MultiModalPlugin.map_input` for more details.
184
185
186

        Note:
            This should be called after :meth:`init_mm_limits_per_prompt`.
187
        """
188
        merged_dict: Dict[str, NestedTensors] = {}
189
190

        for data_key, data_value in data.items():
191
            plugin = self._get_plugin(data_key)
192

193
194
195
196
197
198
199
200
            num_items = len(data_value) if isinstance(data_value, list) else 1
            max_items = self._limits_by_model[model_config][data_key]
            if num_items > max_items:
                raise ValueError(
                    f"You set {data_key}={max_items} (or defaulted to 1) in "
                    f"`--limit-mm-per-prompt`, but found {num_items} items "
                    "in the same prompt.")

201
202
            input_dict = plugin.map_input(model_config, data_value,
                                          mm_processor_kwargs)
203
204
205
206
207
208
209
210
            for input_key, input_tensor in input_dict.items():
                if input_key in merged_dict:
                    raise ValueError(f"The input mappers (keys={set(data)}) "
                                     f"resulted in a conflicting keyword "
                                     f"argument to `forward()`: {input_key}")

                merged_dict[input_key] = input_tensor

211
        return MultiModalKwargs(merged_dict)
212

213
    def create_input_mapper(self, model_config: "ModelConfig"):
214
        """
215
        Create an input mapper (see :meth:`map_input`) for a specific model.
216
        """
217
218
219
220
221
222
223
224
225
        # NOTE - we currently make the assumption that if a model has multiple
        # supported modalities, they take the same kwargs. For the default,
        # this could be an issue in the future if it falls back to two HF
        # resources and we can't inspect the signature easily since it's
        # getting initialized through the autoclass.
        #
        # If this is a problem in the future, we should revisit it, but since
        # it potentially introduces a lot of complexity for a currently
        # uncommon case, we do not for simplicity of both use & implementation
226
        return functools.partial(self.map_input, model_config)
227

228
229
230
231
232
    def register_max_multimodal_tokens(
        self,
        data_type_key: str,
        max_mm_tokens: Optional[MultiModalTokensCalc] = None,
    ):
233
        """
234
235
236
        Register the maximum number of tokens, corresponding to a single
        instance of multimodal data belonging to a specific modality, that are
        passed to the language model for a model class.
237
238
239
240
241
242
243
244
245
        """
        return self._get_plugin(data_type_key) \
            .register_max_multimodal_tokens(max_mm_tokens)

    def register_max_image_tokens(
        self,
        max_mm_tokens: Optional[MultiModalTokensCalc] = None,
    ):
        """
246
247
        Register the maximum number of image tokens, corresponding to a single
        image, that are passed to the language model for a model class.
248
249
250
        """
        return self.register_max_multimodal_tokens("image", max_mm_tokens)

251
252
253
254
255
256
257
258
259
260
261
    def get_max_tokens_per_item_by_modality(
        self,
        model_config: "ModelConfig",
    ) -> Mapping[str, int]:
        """
        Get the maximum number of tokens per data item from each modality
        for profiling the memory usage of a model.

        Note:
            This is currently directly used only in V1.
        """
262
263
264
        if self.has_processor(model_config):
            tokenizer = cached_get_tokenizer(model_config.tokenizer)
            processor = self.create_processor(model_config, tokenizer)
265
            seq_len = model_config.max_model_len
266
            return processor.info.get_mm_max_tokens_per_item(seq_len)
267
268
269
270
271
272

        return {
            key: plugin.get_max_multimodal_tokens(model_config)
            for key, plugin in self._plugins.items()
        }

273
274
275
276
    def get_max_tokens_by_modality(
        self,
        model_config: "ModelConfig",
    ) -> Mapping[str, int]:
277
        """
278
        Get the maximum number of tokens from each modality
279
        for profiling the memory usage of a model.
280

281
        See :meth:`MultiModalPlugin.get_max_multimodal_tokens` for more details.
282
283
284
285
286
287

        Note:
            This should be called after :meth:`init_mm_limits_per_prompt`.
        """
        limits_per_plugin = self._limits_by_model[model_config]

288
        return {
289
290
291
            key: limits_per_plugin[key] * max_tokens_per_mm_item
            for key, max_tokens_per_mm_item in
            self.get_max_tokens_per_item_by_modality(model_config).items()
292
293
294
295
296
297
298
299
300
301
302
303
304
        }

    def get_max_multimodal_tokens(self, model_config: "ModelConfig") -> int:
        """
        Get the maximum number of multi-modal tokens
        for profiling the memory usage of a model.

        See :meth:`MultiModalPlugin.get_max_multimodal_tokens` for more details.

        Note:
            This should be called after :meth:`init_mm_limits_per_prompt`.
        """
        return sum(self.get_max_tokens_by_modality(model_config).values())
305
306
307

    def init_mm_limits_per_prompt(
        self,
308
        model_config: "ModelConfig",
309
310
311
312
313
314
315
316
317
318
    ) -> None:
        """
        Initialize the maximum number of multi-modal input instances for each
        modality that are allowed per prompt for a model class.
        """
        if model_config in self._limits_by_model:
            logger.warning(
                "`mm_limits` has already been set for model=%s, and will "
                "be overwritten by the new values.", model_config.model)

319
        multimodal_config = model_config.multimodal_config
320
321
322
323
324
325
326
327
328
329
330
331
332
333
334
335
336
337
338
339
340
341
342
343
        if multimodal_config is None:
            limits_per_plugin = self._disabled_limits_per_plugin
        else:
            config_limits_per_plugin = multimodal_config.limit_per_prompt

            extra_keys = config_limits_per_plugin.keys() - self._plugins.keys()
            if extra_keys:
                logger.warning(
                    "Detected extra keys in `--limit-mm-per-prompt` which "
                    "are not registered as multi-modal plugins: %s. "
                    "They will be ignored.", extra_keys)

            # NOTE: Currently the default is set to 1 for each plugin
            # TODO: Automatically determine the limits based on budget
            # once more models support multi-image inputs
            limits_per_plugin = {
                key: config_limits_per_plugin.get(key, 1)
                for key in self._plugins
            }

        self._limits_by_model[model_config] = limits_per_plugin

    def get_mm_limits_per_prompt(
        self,
344
        model_config: "ModelConfig",
345
346
347
348
349
350
351
    ) -> Mapping[str, int]:
        """
        Get the maximum number of multi-modal input instances for each modality
        that are allowed per prompt for a model class.

        Note:
            This should be called after :meth:`init_mm_limits_per_prompt`.
352
        """
353
        return self._limits_by_model[model_config]
354
355
356

    def register_processor(
        self,
357
358
359
360
        processor: MultiModalProcessorFactory[_I],
        *,
        info: ProcessingInfoFactory[_I],
        dummy_inputs: DummyInputsBuilderFactory[_I],
361
362
    ):
        """
363
364
        Register a multi-modal processor to a model class. The processor
        is constructed lazily, hence a factory method should be passed.
365
366
367
368
369

        When the model receives multi-modal data, the provided function is
        invoked to transform the data into a dictionary of model inputs.

        See also:
370
371
            - :ref:`input-processing-pipeline`
            - :ref:`enabling-multimodal-inputs`
372
373
374
        """

        def wrapper(model_cls: N) -> N:
375
            if self._processor_factories.contains(model_cls, strict=True):
376
                logger.warning(
377
                    "Model class %s already has a multi-modal processor "
378
379
380
                    "registered to %s. It is overwritten by the new one.",
                    model_cls, self)

381
382
383
384
385
            self._processor_factories[model_cls] = _ProcessorFactories(
                info=info,
                dummy_inputs=dummy_inputs,
                processor=processor,
            )
386
387
388
389
390

            return model_cls

        return wrapper

391
    def _get_model_cls(self, model_config: "ModelConfig"):
392
393
394
395
        # Avoid circular import
        from vllm.model_executor.model_loader import get_model_architecture

        model_cls, _ = get_model_architecture(model_config)
396
397
398
399
400
401
402
        return model_cls

    def has_processor(self, model_config: "ModelConfig") -> bool:
        """
        Test whether a multi-modal processor is defined for a specific model.
        """
        return self._get_model_cls(model_config) in self._processor_factories
403
404
405
406
407

    def create_processor(
        self,
        model_config: "ModelConfig",
        tokenizer: AnyTokenizer,
408
    ) -> BaseMultiModalProcessor[BaseProcessingInfo]:
409
410
411
        """
        Create a multi-modal processor for a specific model and tokenizer.
        """
412
        model_cls = self._get_model_cls(model_config)
413
        factories = self._processor_factories[model_cls]
414
415

        ctx = InputProcessingContext(model_config, tokenizer)
416
417
418
        cache = (None if model_config.disable_mm_preprocessor_cache else
                 self._processing_cache)

419
        return factories.build_processor(ctx, cache=cache)