Skip to content

Explainability/Interpretablity

The explainability features in PyTorch Tabular allow users to interpret and understand the predictions made by a tabular deep learning model. These features provide insights into the model's decision-making process and help identify the most influential features. Some of the explainability features are inbuilt from the models, and a lot of others are based on the Captum library.

Native Feature Importance

One of the features of the GBDT models which everybody loves is the feature importance. It helps us understand which features are the most important for the model. PyTorch Tabular provides a similar feature for some of the models - GANDALF, GATE, and FTTransformers - where the models natively support the extraction of feature importance.

# tabular_model is the trained model of a supported model
tabular_model.feature_importance()

Local Feature Attributions/Explanations

Local feature attributions/explanations help us understand the contribution of each feature towards the prediction for a particular sample. PyTorch Tabular provides this feature for all the models except TabTransformer, Tabnet, and Mixed Density Networks. It is based on the Captum library. The library provides a lot of algorithms for computing feature attributions. PyTorch Tabular provides a wrapper around the library to make it easy to use. The following algorithms are supported:

  • GradientShap: https://captum.ai/api/gradient_shap.html
  • IntegratedGradients: https://captum.ai/api/integrated_gradients.html
  • DeepLift: https://captum.ai/api/deep_lift.html
  • DeepLiftShap: https://captum.ai/api/deep_lift_shap.html
  • InputXGradient: https://captum.ai/api/input_x_gradient.html
  • FeaturePermutation: https://captum.ai/api/feature_permutation.html
  • FeatureAblation: https://captum.ai/api/feature_ablation.html
  • KernelShap: https://captum.ai/api/kernel_shap.html

PyTorch Tabular also supports explaining single instances as well as batches of instances. But, larger datasets will take longer to explain. An exception is the FeaturePermutation and FeatureAblation methods, which is only meaningful for large batches of instances.

Most of these explainability methods require a baseline. This is used to compare the attributions of the input with the attributions of the baseline. The baseline can be a scalar value, a tensor of the same shape as the input, or a special string like "b|10000" which means 10000 samples from the training data. If the baseline is not provided, the default baseline (zero) is used.

# tabular_model is the trained model of a supported model

# Explain a single instance using the GradientShap method and baseline as 10000 samples from the training data
tabular_model.explain(test.head(1), method="GradientShap", baselines="b|10000")

# Explain a batch of instances using the IntegratedGradients method and baseline as 0
tabular_model.explain(test.head(10), method="IntegratedGradients", baselines=0)

Checkout the Captum documentation for more details on the algorithms and the Explainability Tutorial for example usage.

API Reference

pytorch_tabular.TabularModel.explain(data, method='GradientShap', method_args={}, baselines=None, **kwargs)

Returns the feature attributions/explanations of the model as a pandas DataFrame. The shape of the returned dataframe is (num_samples, num_features)

Parameters:

Name Type Description Default
data DataFrame

The dataframe to be explained

required
method str

The method to be used for explaining the model. It should be one of the Defaults to "GradientShap". For more details, refer to https://captum.ai/api/attribution.html

'GradientShap'
method_args Optional[Dict]

The arguments to be passed to the initialization of the Captum method.

{}
baselines Union[float, tensor, str]

The baselines to be used for the explanation. If a scalar is provided, will use that value as the baseline for all the features. If a tensor is provided, will use that tensor as the baseline for all the features. If a string like b|<num_samples> is provided, will use that many samples from the train Using the whole train data as the baseline is not recommended as it can be computationally expensive. By default, PyTorch Tabular uses 10000 samples from the train data as the baseline. You can configure this by passing a special string "b|" where is the number of samples to be used as the baseline. For eg. "b|1000" will use 1000 samples from the train. If None, will use default settings like zero in captum(which is method dependent). For GradientShap, it is the train data. Defaults to None.

None
**kwargs

Additional keyword arguments to be passed to the Captum method attribute function.

{}

Returns:

Name Type Description
DataFrame DataFrame

The dataframe with the feature importance

Source code in src/pytorch_tabular/tabular_model.py
def explain(
    self,
    data: DataFrame,
    method: str = "GradientShap",
    method_args: Optional[Dict] = {},
    baselines: Union[float, torch.tensor, str] = None,
    **kwargs,
) -> DataFrame:
    """Returns the feature attributions/explanations of the model as a pandas DataFrame. The shape of the returned
    dataframe is (num_samples, num_features)

    Args:
        data (DataFrame): The dataframe to be explained
        method (str): The method to be used for explaining the model.
            It should be one of the Defaults to "GradientShap".
            For more details, refer to https://captum.ai/api/attribution.html
        method_args (Optional[Dict], optional): The arguments to be passed to the initialization
            of the Captum method.
        baselines (Union[float, torch.tensor, str]): The baselines to be used for the explanation.
            If a scalar is provided, will use that value as the baseline for all the features.
            If a tensor is provided, will use that tensor as the baseline for all the features.
            If a string like `b|<num_samples>` is provided, will use that many samples from the train
            Using the whole train data as the baseline is not recommended as it can be
            computationally expensive. By default, PyTorch Tabular uses 10000 samples from the
            train data as the baseline. You can configure this by passing a special string
            "b|<num_samples>" where <num_samples> is the number of samples to be used as the
            baseline. For eg. "b|1000" will use 1000 samples from the train.
            If None, will use default settings like zero in captum(which is method dependent).
            For `GradientShap`, it is the train data.
            Defaults to None.

        **kwargs: Additional keyword arguments to be passed to the Captum method `attribute` function.

    Returns:
        DataFrame: The dataframe with the feature importance

    """
    assert CAPTUM_INSTALLED, "Captum not installed. Please install using `pip install captum` or "
    "install PyTorch Tabular using `pip install pytorch-tabular[extra]`"
    ALLOWED_METHODS = [
        "GradientShap",
        "IntegratedGradients",
        "DeepLift",
        "DeepLiftShap",
        "InputXGradient",
        "FeaturePermutation",
        "FeatureAblation",
        "KernelShap",
    ]
    assert method in ALLOWED_METHODS, f"method should be one of {ALLOWED_METHODS}"
    if isinstance(data, pd.Series):
        data = data.to_frame().T
    if method in ["DeepLiftShap", "KernelShap"]:
        warnings.warn(
            f"{method} is computationally expensive and will take some time. For"
            " faster results, try usingsome other methods like GradientShap,"
            " IntegratedGradients etc."
        )
    if method in ["FeaturePermutation", "FeatureAblation"]:
        assert data.shape[0] > 1, f"{method} only works when the number of samples is greater than 1"
        if len(data) <= 100:
            warnings.warn(
                f"{method} gives better results when the number of samples is"
                " large. For better results, try using more samples or some other"
                " methods like GradientShap which works well on single examples."
            )
    is_full_baselines = method in ["GradientShap", "DeepLiftShap"]
    is_not_supported = self.model._get_name() in [
        "TabNetModel",
        "MDNModel",
        "TabTransformerModel",
    ]
    do_baselines = method not in [
        "Saliency",
        "InputXGradient",
        "FeaturePermutation",
        "LRP",
    ]
    if is_full_baselines and (baselines is None or isinstance(baselines, (float, int))):
        raise ValueError(
            f"baselines cannot be a scalar or None for {method}. Please "
            "provide a tensor or a string like `b|<num_samples>`"
        )
    if is_not_supported:
        raise NotImplementedError(f"Attributions are not implemented for {self.model._get_name()}")

    is_embedding1d = isinstance(self.model.embedding_layer, (Embedding1dLayer, PreEncoded1dLayer))
    is_embedding2d = isinstance(self.model.embedding_layer, Embedding2dLayer)
    # Models like NODE may have no embedding dims (doing leaveOneOut encoding) even if categorical_dim > 0
    is_embbeding_dims = (
        hasattr(self.model.hparams, "embedding_dims") and self.model.hparams.embedding_dims is not None
    )
    if (not is_embedding1d) and (not is_embedding2d):
        raise NotImplementedError(
            "Attributions are not implemented for models with this type of" " embedding layer"
        )
    test_dl = self.datamodule.prepare_inference_dataloader(data)
    self.model.eval()
    # prepare import for Captum
    tensor_inp, tensor_tgt = self._prepare_input_for_captum(test_dl)
    baselines = self._prepare_baselines_captum(baselines, test_dl, do_baselines, is_full_baselines)
    # prepare model for Captum
    try:
        interp_model = _CaptumModel(self.model)
        captum_interp_cls = getattr(captum.attr, method)(interp_model, **method_args)
        if do_baselines:
            attributions = captum_interp_cls.attribute(
                tensor_inp,
                baselines=baselines,
                target=(tensor_tgt if self.config.task == "classification" else None),
                **kwargs,
            )
        else:
            attributions = captum_interp_cls.attribute(
                tensor_inp,
                target=(tensor_tgt if self.config.task == "classification" else None),
                **kwargs,
            )
        attributions = self._handle_categorical_embeddings_attributions(
            attributions, is_embedding1d, is_embedding2d, is_embbeding_dims
        )
    finally:
        self.model.train()
    assert attributions.shape[1] == self.model.hparams.continuous_dim + self.model.hparams.categorical_dim, (
        "Something went wrong. The number of features in the attributions"
        f" ({attributions.shape[1]}) does not match the number of features in"
        " the model"
        f" ({self.model.hparams.continuous_dim+self.model.hparams.categorical_dim})"
    )
    return pd.DataFrame(
        attributions.detach().cpu().numpy(),
        columns=self.config.continuous_cols + self.config.categorical_cols,
    )